llama.py 29 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758759760761762763764765766767768769770771772773774775776777778779780781782783784785786787788789790791792793794795796797798799800801802803804805806807808809810811812813814815816817818819820821822823824825826827828829830831832833834835836837838839840841842843844845846847848849850851852853854855856857858859860861862863864865866867868869870871872873874875876877878879880881882883884885886887888889890891892893894895896897898899900901902903
  1. import dataclasses
  2. import json
  3. import math
  4. from collections import OrderedDict
  5. from dataclasses import dataclass
  6. from pathlib import Path
  7. from typing import Optional
  8. import torch
  9. import torch.nn as nn
  10. from einops import rearrange
  11. from loguru import logger
  12. from torch import Tensor
  13. from torch.nn import functional as F
  14. from torch.nn.attention import SDPBackend, sdpa_kernel
  15. from torch.utils.checkpoint import checkpoint
  16. from transformers import AutoTokenizer
  17. from fish_speech.models.text2semantic.lora import LoraConfig, setup_lora
  18. from fish_speech.tokenizer import SEMANTIC_TOKENS, FishTokenizer
  19. def find_multiple(n: int, k: int) -> int:
  20. if n % k == 0:
  21. return n
  22. return n + k - (n % k)
  23. @dataclass
  24. class BaseModelArgs:
  25. model_type: str = "base"
  26. vocab_size: int = 32000
  27. n_layer: int = 32
  28. n_head: int = 32
  29. dim: int = 4096
  30. intermediate_size: int = None
  31. n_local_heads: int = -1
  32. head_dim: int = 64
  33. rope_base: float = 10000
  34. norm_eps: float = 1e-5
  35. max_seq_len: int = 2048
  36. dropout: float = 0.0
  37. tie_word_embeddings: bool = True
  38. attention_qkv_bias: bool = False
  39. attention_o_bias: bool = False
  40. attention_qk_norm: bool = False
  41. # Codebook configs
  42. codebook_size: int = 160
  43. num_codebooks: int = 4
  44. # Gradient checkpointing
  45. use_gradient_checkpointing: bool = True
  46. # Initialize the model
  47. initializer_range: float = 0.02
  48. # Dummy vars
  49. is_reward_model: bool = False
  50. scale_codebook_embeddings: bool = False
  51. def __post_init__(self):
  52. if self.n_local_heads == -1:
  53. self.n_local_heads = self.n_head
  54. if self.intermediate_size is None:
  55. hidden_dim = 4 * self.dim
  56. n_hidden = int(2 * hidden_dim / 3)
  57. self.intermediate_size = find_multiple(n_hidden, 256)
  58. if self.head_dim is None:
  59. self.head_dim = self.dim // self.n_head
  60. @staticmethod
  61. def from_pretrained(path: str):
  62. path = Path(path)
  63. if path.is_dir():
  64. path = path / "config.json"
  65. with open(path, "r", encoding="utf-8") as f:
  66. data = json.load(f)
  67. match data["model_type"]:
  68. case "naive":
  69. cls = NaiveModelArgs
  70. case "dual_ar":
  71. cls = DualARModelArgs
  72. case _:
  73. raise ValueError(f"Unknown model type: {data['model_type']}")
  74. return cls(**data)
  75. def save(self, path: str):
  76. with open(path, "w") as f:
  77. json.dump(self.__dict__, f, indent=4, sort_keys=True, ensure_ascii=False)
  78. @dataclass
  79. class NaiveModelArgs(BaseModelArgs):
  80. model_type: str = "naive"
  81. @dataclass
  82. class DualARModelArgs(BaseModelArgs):
  83. model_type: str = "dual_ar"
  84. n_fast_layer: int = 4
  85. fast_dim: int | None = None
  86. fast_n_head: int | None = None
  87. fast_n_local_heads: int | None = None
  88. fast_head_dim: int | None = None
  89. fast_intermediate_size: int | None = None
  90. fast_attention_qkv_bias: bool | None = None
  91. fast_attention_qk_norm: bool | None = None
  92. fast_attention_o_bias: bool | None = None
  93. def __post_init__(self):
  94. super().__post_init__()
  95. self.fast_dim = self.fast_dim or self.dim
  96. self.fast_n_head = self.fast_n_head or self.n_head
  97. self.fast_n_local_heads = self.fast_n_local_heads or self.n_local_heads
  98. self.fast_head_dim = self.fast_head_dim or self.head_dim
  99. self.fast_intermediate_size = (
  100. self.fast_intermediate_size or self.intermediate_size
  101. )
  102. self.fast_attention_qkv_bias = (
  103. self.fast_attention_qkv_bias
  104. if self.fast_attention_qkv_bias is not None
  105. else self.attention_qkv_bias
  106. )
  107. self.fast_attention_qk_norm = (
  108. self.fast_attention_qk_norm
  109. if self.fast_attention_qk_norm is not None
  110. else self.attention_qk_norm
  111. )
  112. self.fast_attention_o_bias = (
  113. self.fast_attention_o_bias
  114. if self.fast_attention_o_bias is not None
  115. else self.attention_o_bias
  116. )
  117. class KVCache(nn.Module):
  118. def __init__(
  119. self, max_batch_size, max_seq_len, n_heads, head_dim, dtype=torch.bfloat16
  120. ):
  121. super().__init__()
  122. cache_shape = (max_batch_size, n_heads, max_seq_len, head_dim)
  123. self.register_buffer("k_cache", torch.zeros(cache_shape, dtype=dtype))
  124. self.register_buffer("v_cache", torch.zeros(cache_shape, dtype=dtype))
  125. def update(self, input_pos, k_val, v_val):
  126. # input_pos: [S], k_val: [B, H, S, D]
  127. assert input_pos.shape[0] == k_val.shape[2]
  128. k_out = self.k_cache
  129. v_out = self.v_cache
  130. k_out[:, :, input_pos] = k_val
  131. v_out[:, :, input_pos] = v_val
  132. return k_out, v_out
  133. @dataclass
  134. class TransformerForwardResult:
  135. token_logits: Tensor
  136. codebook_logits: Tensor
  137. @dataclass
  138. class BaseTransformerForwardResult:
  139. logits: Tensor
  140. hidden_states: Tensor
  141. class BaseTransformer(nn.Module):
  142. def __init__(
  143. self,
  144. config: BaseModelArgs,
  145. tokenizer: FishTokenizer,
  146. init_weights: bool = True,
  147. ) -> None:
  148. super().__init__()
  149. self.config = config
  150. self.tokenizer = tokenizer
  151. self.semantic_token_ids = list(tokenizer.semantic_id_to_token_id.values())
  152. # Slow transformer
  153. self.embeddings = nn.Embedding(
  154. config.vocab_size,
  155. config.dim,
  156. )
  157. self.codebook_embeddings = nn.Embedding(
  158. config.codebook_size * config.num_codebooks,
  159. config.dim,
  160. )
  161. self.layers = nn.ModuleList(
  162. TransformerBlock(config, use_sdpa=True) for _ in range(config.n_layer)
  163. )
  164. self.norm = RMSNorm(config.dim, eps=config.norm_eps)
  165. if self.config.tie_word_embeddings is False:
  166. self.output = nn.Linear(
  167. config.dim,
  168. config.vocab_size,
  169. bias=False,
  170. )
  171. self.register_buffer(
  172. "freqs_cis",
  173. precompute_freqs_cis(
  174. config.max_seq_len,
  175. config.head_dim,
  176. config.rope_base,
  177. ),
  178. persistent=False,
  179. )
  180. self.register_buffer(
  181. "causal_mask",
  182. torch.tril(
  183. torch.ones(
  184. config.max_seq_len,
  185. config.max_seq_len,
  186. dtype=torch.bool,
  187. )
  188. ),
  189. persistent=False,
  190. )
  191. # For kv cache
  192. self.max_batch_size = -1
  193. self.max_seq_len = -1
  194. if init_weights:
  195. self.apply(self._init_weights)
  196. def setup_caches(
  197. self, max_batch_size: int, max_seq_len: int, dtype: torch.dtype = torch.bfloat16
  198. ):
  199. if self.max_seq_len >= max_seq_len and self.max_batch_size >= max_batch_size:
  200. return
  201. max_seq_len = find_multiple(max_seq_len, 8)
  202. self.max_seq_len = max_seq_len
  203. self.max_batch_size = max_batch_size
  204. for b in self.layers:
  205. b.attention.kv_cache = KVCache(
  206. max_batch_size,
  207. max_seq_len,
  208. self.config.n_local_heads,
  209. self.config.head_dim,
  210. dtype=dtype,
  211. )
  212. def embed(self, inp: Tensor) -> Tensor:
  213. embeds = []
  214. semantic_token_ids_tensor = torch.tensor(
  215. self.semantic_token_ids, device=inp.device, dtype=inp.dtype
  216. )
  217. for i in range(self.config.num_codebooks):
  218. emb = self.codebook_embeddings(
  219. inp[:, i + 1] + i * self.config.codebook_size
  220. )
  221. embeds.append(emb)
  222. vq_embeds_sum = torch.stack(embeds, dim=1).sum(dim=1)
  223. vq_embeds_sum[~torch.isin(inp[:, 0], semantic_token_ids_tensor)] = 0
  224. x = self.embeddings(inp[:, 0]) + vq_embeds_sum
  225. return x
  226. def forward(
  227. self,
  228. inp: Tensor,
  229. key_padding_mask: Optional[Tensor] = None,
  230. ) -> BaseTransformerForwardResult:
  231. seq_len = inp.size(2)
  232. # Here we want to merge the embeddings of the codebooks
  233. x = self.embed(inp)
  234. freqs_cis = self.freqs_cis[:seq_len]
  235. # Not that the causal mask here follows the definition of scaled_dot_product_attention
  236. # That is, FALSE means masked out
  237. # To maintain consistency, key_padding_mask use TRUE to mask out
  238. mask = None
  239. if key_padding_mask is not None:
  240. causal = self.causal_mask[:seq_len, :seq_len]
  241. causal = rearrange(causal, "q k -> 1 1 q k")
  242. atten_mask = rearrange(key_padding_mask, "b s -> b 1 1 s")
  243. atten_mask = atten_mask.logical_not()
  244. mask = causal & atten_mask
  245. # return freqs_cis, mask
  246. for layer in self.layers:
  247. if self.config.use_gradient_checkpointing and self.training:
  248. x = checkpoint(layer, x, freqs_cis, mask, use_reentrant=True)
  249. else:
  250. x = layer(x, freqs_cis, mask)
  251. # We got slow_out here
  252. slow_out = self.norm(x)
  253. if self.config.tie_word_embeddings:
  254. token_logits = F.linear(slow_out, self.embeddings.weight)
  255. else:
  256. token_logits = self.output(slow_out)
  257. return BaseTransformerForwardResult(
  258. logits=token_logits,
  259. hidden_states=x,
  260. )
  261. def forward_generate(
  262. self,
  263. inp: Tensor,
  264. input_pos: Optional[Tensor] = None,
  265. return_all: bool = False,
  266. ) -> BaseTransformerForwardResult:
  267. x = self.embed(inp)
  268. if input_pos is None:
  269. input_pos = torch.arange(inp.shape[-1], device=x.device)
  270. max_seq_len = inp.shape[-1]
  271. else:
  272. max_seq_len = self.max_seq_len
  273. mask = self.causal_mask[None, None, input_pos, :max_seq_len] # (B, N, Q, K)
  274. freqs_cis = self.freqs_cis[input_pos]
  275. for layer in self.layers:
  276. x = layer(x, freqs_cis, mask, input_pos=input_pos)
  277. # If prefill, we only calculate the logits of last token
  278. if x.size(1) > 1 and not return_all:
  279. x = x[:, -1:]
  280. # We got slow_out here
  281. slow_out = self.norm(x)
  282. if self.config.is_reward_model:
  283. token_logits = self.score_output(slow_out)
  284. elif self.config.tie_word_embeddings:
  285. token_logits = F.linear(slow_out, self.embeddings.weight)
  286. else:
  287. token_logits = self.output(slow_out)
  288. return BaseTransformerForwardResult(
  289. logits=token_logits,
  290. hidden_states=x,
  291. )
  292. def _init_weights(self, module):
  293. std = self.config.initializer_range
  294. if isinstance(module, nn.Linear):
  295. module.weight.data.normal_(mean=0.0, std=std)
  296. if module.bias is not None:
  297. module.bias.data.zero_()
  298. elif isinstance(module, nn.Embedding):
  299. module.weight.data.normal_(mean=0.0, std=std)
  300. if module.padding_idx is not None:
  301. module.weight.data[module.padding_idx].zero_()
  302. @staticmethod
  303. def from_pretrained(
  304. path: str,
  305. load_weights: bool = False,
  306. max_length: int | None = None,
  307. lora_config: LoraConfig | None = None,
  308. rope_base: int | None = None,
  309. ) -> "BaseTransformer":
  310. config = BaseModelArgs.from_pretrained(str(path))
  311. if max_length is not None:
  312. config.max_seq_len = max_length
  313. logger.info(f"Override max_seq_len to {max_length}")
  314. if rope_base is not None:
  315. config.rope_base = rope_base
  316. logger.info(f"Override rope_base to {rope_base}")
  317. match config.model_type:
  318. case "naive":
  319. model_cls = NaiveTransformer
  320. case "dual_ar":
  321. model_cls = DualARTransformer
  322. case _:
  323. raise ValueError(f"Unknown model type: {config.model_type}")
  324. tokenizer = FishTokenizer.from_pretrained(path)
  325. logger.info(f"Loading model from {path}, config: {config}")
  326. model = model_cls(config, tokenizer=tokenizer)
  327. if lora_config is not None:
  328. setup_lora(model, lora_config)
  329. logger.info(f"LoRA setup: {lora_config}")
  330. if load_weights is False:
  331. logger.info("Randomly initialized model")
  332. else:
  333. if "int8" in str(Path(path)):
  334. logger.info("Using int8 weight-only quantization!")
  335. from tools.llama.quantize import WeightOnlyInt8QuantHandler
  336. simple_quantizer = WeightOnlyInt8QuantHandler(model)
  337. model = simple_quantizer.convert_for_runtime()
  338. if "int4" in str(Path(path)):
  339. logger.info("Using int4 quantization!")
  340. path_comps = path.name.split("-")
  341. assert path_comps[-2].startswith("g")
  342. groupsize = int(path_comps[-2][1:])
  343. from tools.llama.quantize import WeightOnlyInt4QuantHandler
  344. simple_quantizer = WeightOnlyInt4QuantHandler(model, groupsize)
  345. model = simple_quantizer.convert_for_runtime()
  346. weights = torch.load(
  347. Path(path) / "model.pth",
  348. map_location="cpu",
  349. mmap=True,
  350. weights_only=True,
  351. )
  352. if "state_dict" in weights:
  353. logger.warning(
  354. "Using a TextToSemantic LightningModule checkpoint, "
  355. "please make sure it is a full model, not a LoRA model."
  356. )
  357. weights = weights["state_dict"]
  358. if next(iter(weights.keys())).startswith("model."):
  359. logger.info(
  360. f"Remove prefix 'model.' created by TextToSemantic LightningModule from keys"
  361. )
  362. new_weights = OrderedDict()
  363. for k, v in weights.items():
  364. new_weights[k.replace("model.", "")] = v
  365. weights = new_weights
  366. # Remove audio related weights
  367. for k in list(weights.keys()):
  368. if "audio_" in k:
  369. weights.pop(k)
  370. # Verify the name and shape of parameters since strict=False in load_state_dict.
  371. for k, v in model.named_parameters():
  372. if k not in weights:
  373. logger.warning(f"No weight for {k}")
  374. elif v.shape != weights[k].shape:
  375. logger.warning(
  376. f"Shape mismatch for {k}: {v.shape} vs {weights[k].shape}"
  377. )
  378. err = model.load_state_dict(weights, strict=False, assign=True)
  379. logger.info(f"Loaded weights with error: {err}")
  380. return model
  381. def save_pretrained(self, path: str, drop_lora: bool = False):
  382. path = Path(path)
  383. path.mkdir(parents=True, exist_ok=True)
  384. self.config.save(path / "config.json")
  385. state_dict = self.state_dict()
  386. if drop_lora:
  387. for key in list(state_dict.keys()):
  388. if "lora" not in key:
  389. continue
  390. state_dict.pop(key)
  391. logger.info(f"Drop LoRA parameter: {key}")
  392. torch.save(state_dict, path / "model.pth")
  393. self.tokenizer.save_pretrained(path)
  394. class NaiveTransformer(BaseTransformer):
  395. def __init__(self, config: NaiveModelArgs, tokenizer: FishTokenizer) -> None:
  396. super().__init__(config, init_weights=False, tokenizer=tokenizer)
  397. self.codebook_norm = RMSNorm(config.dim, eps=config.norm_eps)
  398. self.codebook_output = nn.Linear(
  399. config.dim,
  400. config.codebook_size * config.num_codebooks,
  401. bias=False,
  402. )
  403. self.apply(self._init_weights)
  404. def decode(self, result: BaseTransformerForwardResult) -> TransformerForwardResult:
  405. token_logits = result.logits
  406. x = result.hidden_states
  407. # Codebook
  408. codebook_logits = self.codebook_output(self.codebook_norm(x))
  409. codebook_logits = rearrange(
  410. codebook_logits, "b n (c d) -> b n c d", c=self.config.num_codebooks
  411. )
  412. return TransformerForwardResult(
  413. token_logits=token_logits,
  414. codebook_logits=codebook_logits,
  415. )
  416. def forward(
  417. self,
  418. inp: Tensor,
  419. key_padding_mask: Optional[Tensor] = None,
  420. ) -> TransformerForwardResult:
  421. result = super().forward(
  422. inp=inp,
  423. key_padding_mask=key_padding_mask,
  424. )
  425. return self.decode(result)
  426. def forward_generate(
  427. self, x: Tensor, input_pos: Optional[Tensor] = None
  428. ) -> TransformerForwardResult:
  429. result = super().forward_generate(x, input_pos)
  430. return self.decode(result)
  431. class DualARTransformer(BaseTransformer):
  432. def __init__(self, config: NaiveModelArgs, tokenizer: FishTokenizer) -> None:
  433. super().__init__(config, init_weights=False, tokenizer=tokenizer)
  434. # Project to fast dim if needed
  435. if config.fast_dim is not None and config.fast_dim != config.dim:
  436. self.fast_project_in = nn.Linear(config.dim, config.fast_dim)
  437. else:
  438. self.fast_project_in = nn.Identity()
  439. # Fast transformer
  440. self.fast_embeddings = nn.Embedding(config.codebook_size, config.fast_dim)
  441. # The equivalent bs is so large that sdpa doesn't work
  442. override_config = dataclasses.replace(
  443. config,
  444. dim=config.fast_dim,
  445. n_head=config.fast_n_head,
  446. n_local_heads=config.fast_n_local_heads,
  447. head_dim=config.fast_head_dim,
  448. intermediate_size=config.fast_intermediate_size,
  449. attention_qkv_bias=config.fast_attention_qkv_bias,
  450. attention_qk_norm=config.fast_attention_qk_norm,
  451. attention_o_bias=config.fast_attention_o_bias,
  452. )
  453. self.fast_layers = nn.ModuleList(
  454. TransformerBlock(override_config, use_sdpa=False)
  455. for _ in range(config.n_fast_layer)
  456. )
  457. self.fast_norm = RMSNorm(config.fast_dim, eps=config.norm_eps)
  458. self.fast_output = nn.Linear(
  459. config.fast_dim,
  460. config.codebook_size,
  461. bias=False,
  462. )
  463. self.register_buffer(
  464. "fast_freqs_cis",
  465. precompute_freqs_cis(
  466. config.num_codebooks,
  467. config.fast_head_dim,
  468. config.rope_base,
  469. ),
  470. persistent=False,
  471. )
  472. self.apply(self._init_weights)
  473. def setup_caches(
  474. self, max_batch_size: int, max_seq_len: int, dtype: torch.dtype = torch.bfloat16
  475. ):
  476. super().setup_caches(max_batch_size, max_seq_len, dtype)
  477. # Fast transformer
  478. # The max seq len here is the number of codebooks
  479. for b in self.fast_layers:
  480. b.attention.kv_cache = KVCache(
  481. max_batch_size,
  482. self.config.num_codebooks,
  483. self.config.fast_n_local_heads,
  484. self.config.fast_head_dim,
  485. dtype=dtype,
  486. )
  487. def forward(
  488. self,
  489. inp: Tensor,
  490. key_padding_mask: Optional[Tensor] = None,
  491. ) -> TransformerForwardResult:
  492. parent_result = super().forward(inp, key_padding_mask)
  493. token_logits = parent_result.logits
  494. x = parent_result.hidden_states
  495. x = self.fast_project_in(x)
  496. # Fast transformer
  497. fast_seq_len = self.config.num_codebooks
  498. fast_mask = self.causal_mask[
  499. None, None, :fast_seq_len, :fast_seq_len
  500. ] # (B, N, Q, K)
  501. # Drop the last token and rotate left
  502. codebooks = inp[:, 1:-1, 1:]
  503. codebooks = F.pad(codebooks, (0, 1), value=0)
  504. codebook_embeddings = self.fast_embeddings(codebooks)
  505. x = torch.cat([x[:, None], codebook_embeddings], dim=1)
  506. b, s = x.size(0), x.size(2)
  507. x = rearrange(x, "b n s d -> (b s) n d") # flatten the batch and seq_len
  508. # Remove padded part
  509. codebooks = rearrange(codebooks, "b n s -> (b s) n")
  510. codebook_mask = (codebooks == 0).all(dim=-1)
  511. if torch.all(codebook_mask):
  512. # If all codebooks are padded, we keep first 8 to make sure the model runs
  513. codebook_mask[:8] = False
  514. x_bs, x_len = x.size(0), x.size(1)
  515. x = x[~codebook_mask]
  516. for layer in self.fast_layers:
  517. if self.config.use_gradient_checkpointing and self.training:
  518. x = checkpoint(
  519. layer, x, self.fast_freqs_cis, fast_mask, use_reentrant=True
  520. )
  521. else:
  522. x = layer(x, self.fast_freqs_cis, fast_mask)
  523. # unflatten the batch and num_codebooks
  524. fast_out = self.fast_norm(x)
  525. codebook_logits = self.fast_output(fast_out)
  526. # Re-pad the codebook_logits
  527. buffer = torch.zeros(
  528. x_bs,
  529. x_len,
  530. codebook_logits.size(-1),
  531. device=codebook_logits.device,
  532. dtype=codebook_logits.dtype,
  533. )
  534. buffer[~codebook_mask] = codebook_logits
  535. codebook_logits = buffer
  536. assert codebook_logits.shape[1] == self.config.num_codebooks
  537. codebook_logits = rearrange(
  538. codebook_logits,
  539. "(b s) n d -> b s n d",
  540. b=b,
  541. s=s,
  542. n=self.config.num_codebooks,
  543. )
  544. return TransformerForwardResult(
  545. token_logits=token_logits,
  546. codebook_logits=codebook_logits,
  547. )
  548. def forward_generate_fast(
  549. self, x: Tensor, input_pos: Optional[Tensor] = None
  550. ) -> Tensor:
  551. # Fast transformer
  552. x = x.view(1, 1, -1)
  553. fast_mask = self.causal_mask[
  554. None, None, input_pos, : self.config.num_codebooks
  555. ] # (B, N, Q, K)
  556. fast_freqs_cis = self.fast_freqs_cis[input_pos]
  557. for layer in self.fast_layers:
  558. x = layer(x, fast_freqs_cis, fast_mask, input_pos=input_pos)
  559. # unflatten the batch and num_codebooks
  560. fast_out = self.fast_norm(x) # only take the last token
  561. codebook_logits = self.fast_output(fast_out)
  562. return codebook_logits
  563. def forward_generate(
  564. self,
  565. x: Tensor,
  566. input_pos: Optional[Tensor] = None,
  567. vq_masks: Optional[Tensor] = None,
  568. ) -> TransformerForwardResult:
  569. x = super().forward_generate(x, input_pos, vq_masks)
  570. x.hidden_states = self.fast_project_in(x.hidden_states)
  571. return x
  572. class TransformerBlock(nn.Module):
  573. def __init__(self, config: BaseModelArgs, use_sdpa: bool = True) -> None:
  574. super().__init__()
  575. self.attention = Attention(config, use_sdpa=use_sdpa)
  576. self.feed_forward = FeedForward(config)
  577. self.ffn_norm = RMSNorm(config.dim, config.norm_eps)
  578. self.attention_norm = RMSNorm(config.dim, config.norm_eps)
  579. def forward(
  580. self, x: Tensor, freqs_cis: Tensor, mask: Tensor, input_pos: Tensor = None
  581. ) -> Tensor:
  582. h = x + self.attention(self.attention_norm(x), freqs_cis, mask, input_pos)
  583. out = h + self.feed_forward(self.ffn_norm(h))
  584. return out
  585. class Attention(nn.Module):
  586. def __init__(self, config: BaseModelArgs, use_sdpa: bool = True):
  587. super().__init__()
  588. assert config.dim % config.n_head == 0
  589. total_head_dim = (config.n_head + 2 * config.n_local_heads) * config.head_dim
  590. # key, query, value projections for all heads, but in a batch
  591. self.wqkv = nn.Linear(
  592. config.dim, total_head_dim, bias=config.attention_qkv_bias
  593. )
  594. self.wo = nn.Linear(
  595. config.n_head * config.head_dim, config.dim, bias=config.attention_o_bias
  596. )
  597. self.kv_cache = None
  598. if config.attention_qk_norm:
  599. self.q_norm = nn.RMSNorm(config.head_dim, config.norm_eps)
  600. self.k_norm = nn.RMSNorm(config.head_dim, config.norm_eps)
  601. self.dropout = config.dropout
  602. self.n_head = config.n_head
  603. self.head_dim = config.head_dim
  604. self.n_local_heads = config.n_local_heads
  605. self.dim = config.dim
  606. self.use_sdpa = use_sdpa
  607. self.attention_qk_norm = config.attention_qk_norm
  608. self.config = config
  609. self._register_load_state_dict_pre_hook(self.load_hook)
  610. def load_hook(self, state_dict, prefix, *args):
  611. if prefix + "wq.weight" in state_dict:
  612. wq = state_dict.pop(prefix + "wq.weight")
  613. wk = state_dict.pop(prefix + "wk.weight")
  614. wv = state_dict.pop(prefix + "wv.weight")
  615. state_dict[prefix + "wqkv.weight"] = torch.cat([wq, wk, wv])
  616. def forward(
  617. self,
  618. x: Tensor,
  619. freqs_cis: Tensor,
  620. mask: Tensor,
  621. input_pos: Optional[Tensor] = None,
  622. ) -> Tensor:
  623. bsz, seqlen, _ = x.shape
  624. q_size = self.n_head * self.head_dim
  625. kv_size = self.n_local_heads * self.head_dim
  626. q, k, v = self.wqkv(x).split([q_size, kv_size, kv_size], dim=-1)
  627. q = q.view(bsz, seqlen, self.n_head, self.head_dim)
  628. k = k.view(bsz, seqlen, self.n_local_heads, self.head_dim)
  629. v = v.view(bsz, seqlen, self.n_local_heads, self.head_dim)
  630. if self.attention_qk_norm:
  631. q = self.q_norm(q)
  632. k = self.k_norm(k)
  633. q = apply_rotary_emb(q, freqs_cis)
  634. k = apply_rotary_emb(k, freqs_cis)
  635. q, k, v = map(lambda x: x.transpose(1, 2), (q, k, v))
  636. if self.kv_cache is not None:
  637. k, v = self.kv_cache.update(input_pos, k, v)
  638. k = k.repeat_interleave(self.n_head // self.n_local_heads, dim=1)
  639. v = v.repeat_interleave(self.n_head // self.n_local_heads, dim=1)
  640. if self.use_sdpa:
  641. if mask is None:
  642. with sdpa_kernel(SDPBackend.FLASH_ATTENTION):
  643. y = F.scaled_dot_product_attention(
  644. q,
  645. k,
  646. v,
  647. dropout_p=self.dropout if self.training else 0.0,
  648. is_causal=True,
  649. # No third party attn_mask here to use flash_attention
  650. )
  651. else:
  652. y = F.scaled_dot_product_attention(
  653. q,
  654. k,
  655. v,
  656. attn_mask=mask,
  657. dropout_p=self.dropout if self.training else 0.0,
  658. )
  659. else:
  660. y = self.eq_scaled_dot_product_attention(
  661. q,
  662. k,
  663. v,
  664. attn_mask=mask,
  665. dropout_p=self.dropout if self.training else 0.0,
  666. )
  667. y = y.transpose(1, 2).contiguous().view(bsz, seqlen, q_size)
  668. return self.wo(y)
  669. def eq_scaled_dot_product_attention(
  670. self,
  671. query,
  672. key,
  673. value,
  674. attn_mask=None,
  675. dropout_p=0.0,
  676. ) -> torch.Tensor:
  677. # This is a standard scaled dot product attention
  678. # It's low efficient, but it doesn't raise cuda error
  679. L, S = query.size(-2), key.size(-2)
  680. scale_factor = 1 / math.sqrt(query.size(-1))
  681. attn_bias = torch.zeros(1, 1, L, S, dtype=query.dtype, device=query.device)
  682. if attn_mask is not None:
  683. if attn_mask.dtype == torch.bool:
  684. attn_bias.masked_fill_(attn_mask.logical_not(), float("-inf"))
  685. else:
  686. attn_bias += attn_mask
  687. attn_weight = query @ key.transpose(-2, -1) * scale_factor
  688. attn_weight += attn_bias
  689. attn_weight = torch.softmax(attn_weight, dim=-1)
  690. attn_weight = torch.dropout(attn_weight, dropout_p, train=True)
  691. return attn_weight @ value
  692. class FeedForward(nn.Module):
  693. def __init__(self, config: BaseModelArgs) -> None:
  694. super().__init__()
  695. self.w1 = nn.Linear(config.dim, config.intermediate_size, bias=False)
  696. self.w3 = nn.Linear(config.dim, config.intermediate_size, bias=False)
  697. self.w2 = nn.Linear(config.intermediate_size, config.dim, bias=False)
  698. def forward(self, x: Tensor) -> Tensor:
  699. return self.w2(F.silu(self.w1(x)) * self.w3(x))
  700. class RMSNorm(nn.Module):
  701. def __init__(self, dim: int, eps: float = 1e-5):
  702. super().__init__()
  703. self.eps = eps
  704. self.weight = nn.Parameter(torch.ones(dim))
  705. def _norm(self, x):
  706. return x * torch.rsqrt(torch.mean(x * x, dim=-1, keepdim=True) + self.eps)
  707. def forward(self, x: Tensor) -> Tensor:
  708. output = self._norm(x.float()).type_as(x)
  709. return output * self.weight
  710. def precompute_freqs_cis(seq_len: int, n_elem: int, base: int = 10000) -> Tensor:
  711. """
  712. Precomputes frequency tensors for complex exponentials (cis)
  713. Args:
  714. seq_len: Length of the sequence for which positional embeddings are needed.
  715. n_elem: Number of elements in the frequency tensor.
  716. base: Base value for the frequency scaling (default: 10000).
  717. Returns:
  718. A tensor containing the precomputed frequencies in real and imaginary parts (bfloat16).
  719. """
  720. freqs = 1.0 / (
  721. base ** (torch.arange(0, n_elem, 2)[: (n_elem // 2)].float() / n_elem)
  722. )
  723. t = torch.arange(seq_len, device=freqs.device)
  724. freqs = torch.outer(t, freqs)
  725. freqs_cis = torch.polar(torch.ones_like(freqs), freqs)
  726. cache = torch.stack([freqs_cis.real, freqs_cis.imag], dim=-1)
  727. return cache.to(dtype=torch.bfloat16)
  728. def apply_rotary_emb(x: Tensor, freqs_cis: Tensor) -> Tensor:
  729. xshaped = x.float().reshape(*x.shape[:-1], -1, 2)
  730. freqs_cis = freqs_cis.view(1, xshaped.size(1), 1, xshaped.size(3), 2)
  731. x_out2 = torch.stack(
  732. [
  733. xshaped[..., 0] * freqs_cis[..., 0] - xshaped[..., 1] * freqs_cis[..., 1],
  734. xshaped[..., 1] * freqs_cis[..., 0] + xshaped[..., 0] * freqs_cis[..., 1],
  735. ],
  736. -1,
  737. )
  738. x_out2 = x_out2.flatten(3)
  739. return x_out2.type_as(x)