model.py 17 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420
  1. import math
  2. from dataclasses import dataclass
  3. from typing import Tuple, Optional, Literal
  4. import torch
  5. from torch import nn
  6. import torch.nn.functional as F
  7. import torch.distributed as dist
  8. from kernel import act_quant, weight_dequant, fp8_gemm
  9. world_size = 1
  10. rank = 0
  11. block_size = 128
  12. gemm_impl: Literal["bf16", "fp8"] = "bf16"
  13. attn_impl: Literal["naive", "absorb"] = "absorb"
  14. @dataclass
  15. class ModelArgs:
  16. max_batch_size: int = 8
  17. max_seq_len: int = 4096 * 4
  18. dtype: Literal["bf16", "fp8"] = "bf16"
  19. vocab_size: int = 102400
  20. dim: int = 2048
  21. inter_dim: int = 10944
  22. moe_inter_dim: int = 1408
  23. n_layers: int = 27
  24. n_dense_layers: int = 1
  25. n_heads: int = 16
  26. # moe
  27. n_routed_experts: int = 64
  28. n_shared_experts: int = 2
  29. n_activated_experts: int = 6
  30. n_expert_groups: int = 1
  31. n_limited_groups: int = 1
  32. score_func: Literal["softmax", "sigmoid"] = "softmax"
  33. route_scale: float = 1.
  34. # mla
  35. q_lora_rank: int = 0
  36. kv_lora_rank: int = 512
  37. qk_nope_head_dim: int = 128
  38. qk_rope_head_dim: int = 64
  39. v_head_dim: int = 128
  40. # yarn
  41. original_seq_len: int = 4096
  42. rope_theta: float = 10000.0
  43. rope_factor: float = 40
  44. beta_fast: int = 32
  45. beta_slow: int = 1
  46. mscale: float = 1.
  47. class ParallelEmbedding(nn.Module):
  48. def __init__(self, vocab_size: int, dim: int):
  49. super().__init__()
  50. self.vocab_size = vocab_size
  51. self.dim = dim
  52. assert vocab_size % world_size == 0
  53. self.part_vocab_size = (vocab_size // world_size)
  54. self.vocab_start_idx = rank * self.part_vocab_size
  55. self.vocab_end_idx = self.vocab_start_idx + self.part_vocab_size
  56. self.weight = nn.Parameter(torch.empty(self.part_vocab_size, self.dim))
  57. def forward(self, x: torch.Tensor) -> torch.Tensor:
  58. if world_size > 1:
  59. mask = (x < self.vocab_start_idx) | (x >= self.vocab_end_idx)
  60. x = x - self.vocab_start_idx
  61. x[mask] = 0
  62. y = F.embedding(x, self.weight)
  63. if world_size > 1:
  64. y[mask] = 0
  65. dist.all_reduce(y)
  66. return y
  67. def linear(x: torch.Tensor, weight: torch.Tensor, bias: Optional[torch.Tensor] = None) -> torch.Tensor:
  68. if weight.element_size() > 1:
  69. return F.linear(x, weight, bias)
  70. elif gemm_impl == "bf16":
  71. weight = weight_dequant(weight, weight.scale)
  72. return F.linear(x, weight, bias)
  73. else:
  74. x, scale = act_quant(x, block_size)
  75. y = fp8_gemm(x, scale, weight, weight.scale)
  76. if bias is not None:
  77. y += bias
  78. return y
  79. class Linear(nn.Module):
  80. dtype = torch.bfloat16
  81. def __init__(self, in_features: int, out_features: int, bias: bool = False, dtype = None):
  82. super().__init__()
  83. self.in_features = in_features
  84. self.out_features = out_features
  85. self.weight = nn.Parameter(torch.empty(out_features, in_features, dtype=dtype or Linear.dtype))
  86. if self.weight.element_size() == 1:
  87. scale_out_features = (out_features + block_size - 1) // block_size
  88. scale_in_features = (in_features + block_size - 1) // block_size
  89. self.weight.scale = self.scale = nn.Parameter(torch.empty(scale_out_features, scale_in_features, dtype=torch.float32))
  90. else:
  91. self.register_parameter("scale", None)
  92. if bias:
  93. self.bias = nn.Parameter(torch.empty(self.part_out_features))
  94. else:
  95. self.register_parameter("bias", None)
  96. def forward(self, x: torch.Tensor) -> torch.Tensor:
  97. return linear(x, self.weight, self.bias)
  98. class ColumnParallelLinear(Linear):
  99. def __init__(self, in_features: int, out_features: int, bias: bool = False, dtype = None):
  100. assert out_features % world_size == 0
  101. self.part_out_features = out_features // world_size
  102. super().__init__(in_features, self.part_out_features, bias, dtype)
  103. def forward(self, x: torch.Tensor) -> torch.Tensor:
  104. y = linear(x, self.weight, self.bias)
  105. return y
  106. class RowParallelLinear(Linear):
  107. def __init__(self, in_features: int, out_features: int, bias: bool = False, dtype = None):
  108. assert in_features % world_size == 0
  109. self.part_in_features = in_features // world_size
  110. super().__init__(self.part_in_features, out_features, bias, dtype)
  111. def forward(self, x: torch.Tensor) -> torch.Tensor:
  112. y = linear(x, self.weight)
  113. if world_size > 1:
  114. dist.all_reduce(y)
  115. if self.bias is not None:
  116. y += self.bias
  117. return y
  118. class RMSNorm(nn.Module):
  119. def __init__(self, dim: int, eps: float = 1e-6):
  120. super().__init__()
  121. self.dim = dim
  122. self.eps = eps
  123. self.weight = nn.Parameter(torch.ones(dim))
  124. def forward(self, x: torch.Tensor):
  125. return F.rms_norm(x, (self.dim,), self.weight, self.eps)
  126. def precompute_freqs_cis(args: ModelArgs) -> torch.Tensor:
  127. dim = args.qk_rope_head_dim
  128. seqlen = args.max_seq_len
  129. beta_fast = args.beta_fast
  130. beta_slow = args.beta_slow
  131. base = args.rope_theta
  132. factor = args.rope_factor
  133. def find_correction_dim(num_rotations, dim, base, max_seq_len):
  134. return dim * math.log(max_seq_len / (num_rotations * 2 * math.pi)) / (2 * math.log(base))
  135. def find_correction_range(low_rot, high_rot, dim, base, max_seq_len):
  136. low = math.floor(find_correction_dim(low_rot, dim, base, max_seq_len))
  137. high = math.ceil(find_correction_dim(high_rot, dim, base, max_seq_len))
  138. return max(low, 0), min(high, dim-1)
  139. def linear_ramp_factor(min, max, dim):
  140. if min == max:
  141. max += 0.001
  142. linear_func = (torch.arange(dim, dtype=torch.float32) - min) / (max - min)
  143. ramp_func = torch.clamp(linear_func, 0, 1)
  144. return ramp_func
  145. freqs = 1.0 / (base ** (torch.arange(0, dim, 2, dtype=torch.float32) / dim))
  146. if seqlen > args.original_seq_len:
  147. low, high = find_correction_range(beta_fast, beta_slow, dim, base, args.original_seq_len)
  148. smooth = 1 - linear_ramp_factor(low, high, dim // 2)
  149. freqs = freqs / factor * (1 - smooth) + freqs * smooth
  150. t = torch.arange(seqlen)
  151. freqs = torch.outer(t, freqs)
  152. freqs_cis = torch.polar(torch.ones_like(freqs), freqs)
  153. return freqs_cis
  154. def apply_rotary_emb(x: torch.Tensor, freqs_cis: torch.Tensor) -> torch.Tensor:
  155. dtype = x.dtype
  156. x = torch.view_as_complex(x.float().view(*x.shape[:-1], -1, 2))
  157. freqs_cis = freqs_cis.view(1, x.size(1), 1, x.size(-1))
  158. y = torch.view_as_real(x * freqs_cis).flatten(3)
  159. return y.to(dtype)
  160. class MLA(nn.Module):
  161. def __init__(self, args: ModelArgs):
  162. super().__init__()
  163. self.dim = args.dim
  164. self.n_heads = args.n_heads
  165. self.n_local_heads = args.n_heads // world_size
  166. self.q_lora_rank = args.q_lora_rank
  167. self.kv_lora_rank = args.kv_lora_rank
  168. self.qk_nope_head_dim = args.qk_nope_head_dim
  169. self.qk_rope_head_dim = args.qk_rope_head_dim
  170. self.qk_head_dim = args.qk_nope_head_dim + args.qk_rope_head_dim
  171. self.v_head_dim = args.v_head_dim
  172. if self.q_lora_rank == 0:
  173. self.wq = ColumnParallelLinear(self.dim, self.n_heads * self.qk_head_dim)
  174. else:
  175. self.wq_a = Linear(self.dim, self.q_lora_rank)
  176. self.q_norm = RMSNorm(self.q_lora_rank)
  177. self.wq_b = ColumnParallelLinear(self.q_lora_rank, self.n_heads * self.qk_head_dim)
  178. self.wkv_a = Linear(self.dim, self.kv_lora_rank + self.qk_rope_head_dim)
  179. self.kv_norm = RMSNorm(self.kv_lora_rank)
  180. self.wkv_b = ColumnParallelLinear(self.kv_lora_rank, self.n_heads * (self.qk_nope_head_dim + self.v_head_dim))
  181. self.wo = RowParallelLinear(self.n_heads * self.v_head_dim, self.dim)
  182. self.softmax_scale = self.qk_head_dim ** -0.5
  183. if args.max_seq_len > args.original_seq_len:
  184. mscale = 0.1 * args.mscale * math.log(args.rope_factor) + 1.0
  185. self.softmax_scale = self.softmax_scale * mscale * mscale
  186. if attn_impl == "naive":
  187. self.register_buffer("k_cache", torch.zeros(args.max_batch_size, args.max_seq_len, self.n_local_heads, self.qk_head_dim), persistent=False)
  188. self.register_buffer("v_cache", torch.zeros(args.max_batch_size, args.max_seq_len, self.n_local_heads, self.v_head_dim), persistent=False)
  189. else:
  190. self.register_buffer("kv_cache", torch.zeros(args.max_batch_size, args.max_seq_len, self.kv_lora_rank), persistent=False)
  191. self.register_buffer("pe_cache", torch.zeros(args.max_batch_size, args.max_seq_len, self.qk_rope_head_dim), persistent=False)
  192. def forward(self, x: torch.Tensor, start_pos: int, freqs_cis: torch.Tensor, mask: Optional[torch.Tensor]):
  193. bsz, seqlen, _ = x.size()
  194. end_pos = start_pos + seqlen
  195. if self.q_lora_rank == 0:
  196. q = self.wq(x)
  197. else:
  198. q = self.wq_b(self.q_norm(self.wq_a(x)))
  199. q = q.view(bsz, seqlen, self.n_local_heads, self.qk_head_dim)
  200. q_nope, q_pe = torch.split(q, [self.qk_nope_head_dim, self.qk_rope_head_dim], dim=-1)
  201. q_pe = apply_rotary_emb(q_pe, freqs_cis)
  202. kv = self.wkv_a(x)
  203. kv, k_pe = torch.split(kv, [self.kv_lora_rank, self.qk_rope_head_dim], dim=-1)
  204. k_pe = apply_rotary_emb(k_pe.unsqueeze(2), freqs_cis)
  205. if attn_impl == "naive":
  206. q = torch.cat([q_nope, q_pe], dim=-1)
  207. kv = self.wkv_b(self.kv_norm(kv))
  208. kv = kv.view(bsz, seqlen, self.n_local_heads, self.qk_nope_head_dim + self.v_head_dim)
  209. k_nope, v = torch.split(kv, [self.qk_nope_head_dim, self.v_head_dim], dim=-1)
  210. k = torch.cat([k_nope, k_pe.expand(-1, -1, self.n_local_heads, -1)], dim=-1)
  211. self.k_cache[:bsz, start_pos:end_pos] = k
  212. self.v_cache[:bsz, start_pos:end_pos] = v
  213. scores = torch.einsum("bshd,bthd->bsht", q, self.k_cache[:bsz, :end_pos]) * self.softmax_scale
  214. else:
  215. wkv_b = self.wkv_b.weight if self.wkv_b.scale is None else weight_dequant(self.wkv_b.weight, self.wkv_b.scale, block_size)
  216. wkv_b = wkv_b.view(self.n_local_heads, -1, self.kv_lora_rank)
  217. q_nope = torch.einsum("bshd,hdc->bshc", q_nope, wkv_b[:, :self.qk_nope_head_dim])
  218. self.kv_cache[:bsz, start_pos:end_pos] = self.kv_norm(kv)
  219. self.pe_cache[:bsz, start_pos:end_pos] = k_pe.squeeze(2)
  220. scores = (torch.einsum("bshc,btc->bsht", q_nope, self.kv_cache[:bsz, :end_pos]) +
  221. torch.einsum("bshr,btr->bsht", q_pe, self.pe_cache[:bsz, :end_pos])) * self.softmax_scale
  222. if mask is not None:
  223. scores += mask.unsqueeze(1)
  224. scores = scores.softmax(dim=-1, dtype=torch.float32).type_as(x)
  225. if attn_impl == "naive":
  226. x = torch.einsum("bsht,bthd->bshd", scores, self.v_cache[:bsz, :end_pos])
  227. else:
  228. x = torch.einsum("bsht,btc->bshc", scores, self.kv_cache[:bsz, :end_pos])
  229. x = torch.einsum("bshc,hdc->bshd", x, wkv_b[:, -self.v_head_dim:])
  230. x = self.wo(x.flatten(2))
  231. return x
  232. class MLP(nn.Module):
  233. def __init__(self, dim: int, inter_dim: int):
  234. super().__init__()
  235. self.w1 = ColumnParallelLinear(dim, inter_dim)
  236. self.w2 = RowParallelLinear(inter_dim, dim)
  237. self.w3 = ColumnParallelLinear(dim, inter_dim)
  238. def forward(self, x: torch.Tensor) -> torch.Tensor:
  239. return self.w2(F.silu(self.w1(x)) * self.w3(x))
  240. class Gate(nn.Module):
  241. def __init__(self, args: ModelArgs):
  242. super().__init__()
  243. self.dim = args.dim
  244. self.topk = args.n_activated_experts
  245. self.n_groups = args.n_expert_groups
  246. self.topk_groups = args.n_limited_groups
  247. self.score_func = args.score_func
  248. self.route_scale = args.route_scale
  249. self.weight = nn.Parameter(torch.empty(args.n_routed_experts, args.dim))
  250. self.bias = nn.Parameter(torch.empty(args.n_routed_experts)) if self.dim == 7168 else None
  251. def forward(self, x: torch.Tensor) -> Tuple[torch.Tensor, torch.Tensor]:
  252. scores = linear(x, self.weight)
  253. if self.score_func == "softmax":
  254. scores = scores.softmax(dim=-1, dtype=torch.float32)
  255. else:
  256. scores = scores.sigmoid()
  257. original_scores = scores
  258. if self.bias is not None:
  259. scores = scores + self.bias
  260. if self.n_groups > 1:
  261. scores = scores.view(x.size(0), self.n_groups, -1)
  262. if self.bias is None:
  263. group_scores = scores.amax(dim=-1)
  264. else:
  265. group_scores = scores.topk(2, dim=-1)[0].sum(dim=-1)
  266. indices = group_scores.topk(self.topk_groups, dim=-1)[1]
  267. mask = torch.zeros_like(scores[..., 0]).scatter_(1, indices, True)
  268. scores = (scores * mask.unsqueeze(-1)).flatten(1)
  269. indices = torch.topk(scores, self.topk, dim=-1)[1]
  270. weights = original_scores.gather(1, indices)
  271. if self.score_func == "sigmoid":
  272. weights /= weights.sum(dim=-1, keepdim=True)
  273. weights *= self.route_scale
  274. return weights.type_as(x), indices
  275. class Expert(nn.Module):
  276. def __init__(self, dim: int, inter_dim: int):
  277. super().__init__()
  278. self.w1 = Linear(dim, inter_dim)
  279. self.w2 = Linear(inter_dim, dim)
  280. self.w3 = Linear(dim, inter_dim)
  281. def forward(self, x: torch.Tensor) -> torch.Tensor:
  282. return self.w2(F.silu(self.w1(x)) * self.w3(x))
  283. class MoE(nn.Module):
  284. def __init__(self, args: ModelArgs):
  285. super().__init__()
  286. self.dim = args.dim
  287. assert args.n_routed_experts % world_size == 0
  288. self.n_routed_experts = args.n_routed_experts
  289. self.n_local_experts = args.n_routed_experts // world_size
  290. self.n_activated_experts = args.n_activated_experts
  291. self.experts_start_idx = rank * self.n_local_experts
  292. self.experts_end_idx = self.experts_start_idx + self.n_local_experts
  293. self.gate = Gate(args)
  294. self.experts = nn.ModuleList([Expert(args.dim, args.moe_inter_dim) if self.experts_start_idx <= i < self.experts_end_idx else None
  295. for i in range(self.n_routed_experts)])
  296. self.shared_experts = MLP(args.dim, args.n_shared_experts * args.moe_inter_dim)
  297. def forward(self, x: torch.Tensor) -> torch.Tensor:
  298. shape = x.size()
  299. x = x.view(-1, self.dim)
  300. weights, indices = self.gate(x)
  301. y = torch.zeros_like(x)
  302. counts = torch.bincount(indices.flatten(), minlength=self.n_routed_experts).tolist()
  303. for i in range(self.experts_start_idx, self.experts_end_idx):
  304. if counts[i] == 0:
  305. continue
  306. expert = self.experts[i]
  307. idx, top = torch.where(indices == i)
  308. y[idx] += expert(x[idx]) * weights[idx, top, None]
  309. z = self.shared_experts(x)
  310. if world_size > 1:
  311. dist.all_reduce(y)
  312. return (y + z).view(shape)
  313. class Block(nn.Module):
  314. def __init__(self, layer_id: int, args: ModelArgs):
  315. super().__init__()
  316. self.attn = MLA(args)
  317. self.ffn = MLP(args.dim, args.inter_dim) if layer_id < args.n_dense_layers else MoE(args)
  318. self.attn_norm = RMSNorm(args.dim)
  319. self.ffn_norm = RMSNorm(args.dim)
  320. def forward(self, x: torch.Tensor, start_pos: int, freqs_cis: torch.Tensor, mask: Optional[torch.Tensor]) -> torch.Tensor:
  321. x = x + self.attn(self.attn_norm(x), start_pos, freqs_cis, mask)
  322. x = x + self.ffn(self.ffn_norm(x))
  323. return x
  324. class Transformer(nn.Module):
  325. def __init__(self, args: ModelArgs):
  326. global world_size, rank
  327. world_size = dist.get_world_size() if dist.is_initialized() else 1
  328. rank = dist.get_rank() if dist.is_initialized() else 0
  329. Linear.dtype = torch.float8_e4m3fn if args.dtype == "fp8" else torch.bfloat16
  330. super().__init__()
  331. self.max_seq_len = args.max_seq_len
  332. self.embed = ParallelEmbedding(args.vocab_size, args.dim)
  333. self.layers = torch.nn.ModuleList()
  334. for layer_id in range(args.n_layers):
  335. self.layers.append(Block(layer_id, args))
  336. self.norm = RMSNorm(args.dim)
  337. self.head = ColumnParallelLinear(args.dim, args.vocab_size, dtype=torch.get_default_dtype())
  338. self.register_buffer("freqs_cis", precompute_freqs_cis(args), persistent=False)
  339. @torch.inference_mode()
  340. def forward(self, tokens: torch.Tensor, start_pos: int = 0):
  341. seqlen = tokens.size(1)
  342. h = self.embed(tokens)
  343. freqs_cis = self.freqs_cis[start_pos:start_pos+seqlen]
  344. mask = None
  345. if seqlen > 1:
  346. mask = torch.full((seqlen, seqlen), float("-inf"), device=tokens.device).triu_(1)
  347. for layer in self.layers:
  348. h = layer(h, start_pos, freqs_cis, mask)
  349. h = self.norm(h)[:, -1]
  350. logits = self.head(h)
  351. if world_size > 1:
  352. all_logits = [torch.empty_like(logits) for _ in range(world_size)]
  353. dist.all_gather(all_logits, logits)
  354. logits = torch.cat(all_logits, dim=-1)
  355. return logits
  356. if __name__ == "__main__":
  357. torch.set_default_dtype(torch.bfloat16)
  358. torch.set_default_device("cuda")
  359. torch.manual_seed(0)
  360. args = ModelArgs()
  361. x = torch.randint(0, args.vocab_size, (2, 128))
  362. model = Transformer(args)
  363. print(model(x).size())