llama2.py 6.2 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158
  1. # Copyright (c) Microsoft Corporation.
  2. # SPDX-License-Identifier: Apache-2.0
  3. # DeepSpeed Team
  4. from .base import *
  5. from .features import HybridSplitQKVContainer, HybridGatedMLPContainer, MetaTensorContainer
  6. from deepspeed.utils.types import ActivationFuncType, NormType
  7. from deepspeed.model_implementations.transformers.ds_llama2 import DeepSpeedLlama2Inference
  8. import torch
  9. from torch.nn.parameter import Parameter
  10. from ..policy import (
  11. TransformerPolicy,
  12. transformer_param_names,
  13. maybe_copy,
  14. maybe_copy_qkv,
  15. maybe_copy_geglu,
  16. maybe_get_lora,
  17. )
  18. class DS_LLAMA2Container(MetaTensorContainer, HybridGatedMLPContainer, HybridSplitQKVContainer,
  19. BaseTransformerContainer):
  20. def __init__(self, **kwargs):
  21. super().__init__(**kwargs)
  22. # All model specific things should be defined here instead of the base class.
  23. def create_module(self, config=None):
  24. _config = config if config is not None else self.ds_model_config
  25. _config.rotate_half = False
  26. _config.rotate_every_two = True
  27. _config.rotary_dim = self.hidden_size // self.num_attention_heads
  28. _config.num_kv = self.policy.client_module.attention.n_kv_heads
  29. self.module = DeepSpeedLlama2Inference(_config, mp_group=self.mp_group)
  30. return self.module
  31. def set_lora_params(self):
  32. """
  33. Necessary to implement for `HybridEngineContainer`
  34. """
  35. self.lora_params = [
  36. maybe_get_lora(p) for p in [
  37. self.policy.client_module.feed_forward.w3.weight, self.policy.client_module.feed_forward.w1.weight,
  38. self.policy.client_module.feed_forward.w2.weight, self.policy.client_module.attention.wq.weight,
  39. self.policy.client_module.attention.wk.weight, self.policy.client_module.attention.wv.weight,
  40. self.policy.client_module.attention.wo.weight
  41. ]
  42. ]
  43. def get_lora_matched_pair(self):
  44. up_proj_lora, gate_proj_lora, down_proj_lora, q_lora, k_lora, v_lora, out_lora = self.get_lora_params()
  45. ret = [(up_proj_lora, self.inter_up_w), (gate_proj_lora, self.inter_gate_w), (down_proj_lora, self._4hh_w),
  46. (out_lora, self.dense_w), (q_lora, self.qw), (k_lora, self.kw), (v_lora, self.vw)]
  47. return ret
  48. def set_q_k_v(self):
  49. """
  50. Necessary to implement for `HybridSplitQKVContainer`
  51. """
  52. self.qw = self.policy.client_module.attention.wq.weight
  53. self.qb = None
  54. self.kw = self.policy.client_module.attention.wk.weight
  55. self.kb = None
  56. self.vw = self.policy.client_module.attention.wv.weight
  57. self.vb = None
  58. def set_mlp_gate(self):
  59. """
  60. Necessary to implement for `HybridGatedMLPContainer`
  61. """
  62. self.inter_up_w = self.policy.client_module.feed_forward.w2.weight
  63. self.inter_up_b = None
  64. self.inter_gate_w = self.policy.client_module.feed_forward.w1.weight
  65. self.inter_gate_b = None
  66. def load_params(self, module, sd, weight_quantizer, mp_replace, prefix):
  67. param_names = (
  68. 'attention.wq.weight', \
  69. 'attention.wk.weight', \
  70. 'attention.wv.weight', \
  71. 'attention.wo.weight', \
  72. 'feed_forward.w3.weight', \
  73. 'feed_forward.w1.weight', \
  74. 'feed_forward.w2.weight', \
  75. 'ffn_norm.weight', \
  76. 'attention_norm.weight'
  77. )
  78. maybe_copy_qkv(module.attention,
  79. sd,
  80. weight_quantizer,
  81. mp_replace,
  82. 'attn_qkvw', [prefix + param_names[0], prefix + param_names[1], prefix + param_names[2]],
  83. split_qkv=self.policy.split_qkv)
  84. for i in range(3, 4):
  85. maybe_copy(module.attention, sd, weight_quantizer, mp_replace, transformer_param_names[i - 1],
  86. prefix + param_names[i])
  87. maybe_copy_geglu(module.mlp, sd, weight_quantizer, mp_replace, 'inter_w',
  88. [prefix + param_names[4], prefix + param_names[5]])
  89. maybe_copy(module.mlp, sd, weight_quantizer, mp_replace, 'output_w', prefix + param_names[6])
  90. maybe_copy(module.mlp, sd, weight_quantizer, mp_replace, transformer_param_names[8], prefix + param_names[7])
  91. maybe_copy(module, sd, weight_quantizer, mp_replace, transformer_param_names[10], prefix + param_names[8])
  92. class LLAMA2LayerPolicy(TransformerPolicy):
  93. def __init__(self, client_module, inference=True):
  94. super().__init__(
  95. inference,
  96. mlp_act_func_type=ActivationFuncType.GATED_SILU,
  97. norm_type=NormType.RMSNorm,
  98. )
  99. self.client_module = client_module
  100. try:
  101. import llama
  102. LLAMA2LayerPolicy._orig_layer_class = llama.model.TransformerBlock # type: ignore
  103. except:
  104. LLAMA2LayerPolicy._orig_layer_class = None
  105. def get_hidden_heads(self):
  106. return self.client_module.attention.wq.weight.shape[1], \
  107. self.client_module.n_heads, \
  108. self.client_module.ffn_norm.eps, \
  109. (self.client_module.feed_forward.w1.weight.shape[0] * \
  110. deepspeed.comm.get_world_size() if deepspeed.comm.is_initialized() else 1) # this is a hack to inject when model is already partitioned!
  111. def attention(self, enable_training=False):
  112. qw = self.client_module.attention.wq.weight
  113. kw = self.client_module.attention.wk.weight
  114. vw = self.client_module.attention.wv.weight
  115. qkvw = Parameter(torch.cat((qw, kw, vw), dim=0), requires_grad=enable_training)
  116. return qkvw, \
  117. None, \
  118. self.client_module.attention.wo.weight, \
  119. None
  120. def mlp(self, enable_training=False):
  121. mlp1_up = self.client_module.feed_forward.w3.weight
  122. mlp1_gate = self.client_module.feed_forward.w1.weight
  123. mlp2 = self.client_module.feed_forward.w2.weight
  124. mlp1 = Parameter(torch.cat((mlp1_up, mlp1_gate), dim=0), requires_grad=enable_training)
  125. return mlp1, None, mlp2, None
  126. def layernorm(self):
  127. return self.client_module.ffn_norm.weight, \
  128. None, \
  129. self.client_module.attention_norm.weight, \
  130. None