lit_module.py 11 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328
  1. from dataclasses import dataclass
  2. from typing import Any, Optional
  3. import lightning as L
  4. import loralib as lora
  5. import torch
  6. import torch.nn.functional as F
  7. from lightning.pytorch.utilities.types import OptimizerLRScheduler
  8. import fish_speech.utils as utils
  9. from fish_speech.models.text2semantic.llama import NaiveTransformer
  10. log = utils.RankedLogger(__name__, rank_zero_only=True)
  11. @dataclass
  12. class LoraConfig:
  13. r: int
  14. lora_alpha: float
  15. lora_dropout: float = 0.0
  16. class TextToSemantic(L.LightningModule):
  17. def __init__(
  18. self,
  19. model: NaiveTransformer,
  20. optimizer: Any,
  21. lr_scheduler: Any,
  22. lora_config: Optional[LoraConfig] = None,
  23. save_lora_only: bool = False,
  24. use_dpo: bool = False,
  25. dpo_beta: float = 0.2,
  26. ):
  27. super().__init__()
  28. self.model = model
  29. self.optimizer_builder = optimizer
  30. self.lr_scheduler_builder = lr_scheduler
  31. self.lora_config = lora_config
  32. self.save_lora_only = save_lora_only
  33. self.use_dpo = use_dpo # We don't support reference model yet
  34. self.dpo_beta = dpo_beta
  35. if self.lora_config is not None:
  36. self.setup_lora()
  37. def setup_lora(self):
  38. # Replace the embedding layer with a LoRA layer
  39. self.model.embeddings = lora.Embedding(
  40. num_embeddings=self.model.embeddings.num_embeddings,
  41. embedding_dim=self.model.embeddings.embedding_dim,
  42. padding_idx=self.model.embeddings.padding_idx,
  43. r=self.lora_config.r,
  44. lora_alpha=self.lora_config.lora_alpha,
  45. )
  46. # Replace output layer with a LoRA layer
  47. linears = [(self.model, "output")]
  48. # Replace all linear layers with LoRA layers
  49. for layer in self.model.layers:
  50. linears.extend([(layer.attention, "wqkv"), (layer.attention, "wo")])
  51. linears.extend(
  52. [
  53. (layer.feed_forward, "w1"),
  54. (layer.feed_forward, "w2"),
  55. (layer.feed_forward, "w3"),
  56. ]
  57. )
  58. for module, layer in linears:
  59. updated_linear = lora.Linear(
  60. in_features=getattr(module, layer).in_features,
  61. out_features=getattr(module, layer).out_features,
  62. bias=getattr(module, layer).bias,
  63. r=self.lora_config.r,
  64. lora_alpha=self.lora_config.lora_alpha,
  65. lora_dropout=self.lora_config.lora_dropout,
  66. )
  67. setattr(module, layer, updated_linear)
  68. # Mark only the LoRA layers as trainable
  69. lora.mark_only_lora_as_trainable(self.model, bias="lora_only")
  70. def forward(self, x):
  71. return self.model(x)
  72. def on_save_checkpoint(self, checkpoint):
  73. if self.lora_config is None or self.save_lora_only is False:
  74. return
  75. # Save only LoRA parameters
  76. state_dict = checkpoint["state_dict"]
  77. for name in list(state_dict.keys()):
  78. if "lora" not in name:
  79. state_dict.pop(name)
  80. def configure_optimizers(self) -> OptimizerLRScheduler:
  81. # Get weight decay parameters
  82. weight_decay_parameters, other_parameters = [], []
  83. for name, param in self.named_parameters():
  84. if ".bias" in name or "norm.weight" in name or ".embeddings." in name:
  85. other_parameters.append(param)
  86. else:
  87. weight_decay_parameters.append(param)
  88. optimizer = self.optimizer_builder(
  89. [
  90. {"params": weight_decay_parameters},
  91. {"params": other_parameters, "weight_decay": 0.0},
  92. ]
  93. )
  94. # Print the parameters and their weight decay
  95. for i in optimizer.param_groups:
  96. log.info(
  97. f"Set weight decay: {i['weight_decay']} for {len(i['params'])} parameters"
  98. )
  99. lr_scheduler = self.lr_scheduler_builder(optimizer)
  100. return {
  101. "optimizer": optimizer,
  102. "lr_scheduler": {
  103. "scheduler": lr_scheduler,
  104. "interval": "step",
  105. },
  106. }
  107. # Copied from https://github.com/eric-mitchell/direct-preference-optimization/blob/main/trainers.py#L90
  108. def get_batch_logps(
  109. self,
  110. logits: torch.FloatTensor,
  111. labels: torch.LongTensor,
  112. average_log_prob: bool = False,
  113. ) -> torch.FloatTensor:
  114. """Compute the log probabilities of the given labels under the given logits.
  115. Args:
  116. logits: Logits of the model (unnormalized). Shape: (batch_size, sequence_length, codebook_size, vocab_size)
  117. labels: Labels for which to compute the log probabilities. Label tokens with a value of -100 are ignored. Shape: (batch_size, sequence_length, codebook_size)
  118. average_log_prob: If True, return the average log probability per (non-masked) token. Otherwise, return the sum of the log probabilities of the (non-masked) tokens.
  119. Returns:
  120. A tensor of shape (batch_size,) containing the average/sum log probabilities of the given labels under the given logits.
  121. """
  122. assert logits.shape[:-1] == labels.shape
  123. labels = labels.clone()
  124. loss_mask = labels != -100
  125. # dummy token; we'll ignore the losses on these tokens later
  126. labels[labels == -100] = 0
  127. per_token_logps = torch.gather(
  128. logits.log_softmax(-1), dim=-1, index=labels.unsqueeze(-1)
  129. ).squeeze(-1)
  130. if average_log_prob:
  131. return (per_token_logps * loss_mask).sum(-1) / loss_mask.sum(-1)
  132. else:
  133. return (per_token_logps * loss_mask).sum(-1)
  134. def _step(self, batch, batch_idx, stage: str):
  135. # Do positive and negative samples in the same batch to speed up training
  136. labels = batch["labels"]
  137. outputs = self.model(
  138. inp=batch["inputs"],
  139. key_padding_mask=batch["attention_masks"],
  140. )
  141. token_logits = outputs.token_logits
  142. codebook_logits = outputs.codebook_logits
  143. if self.use_dpo:
  144. # Firtst half is positive, second half is negative
  145. token_logits, negative_token_logits = token_logits.chunk(2)
  146. codebook_logits, negative_codebook_logits = codebook_logits.chunk(2)
  147. labels, negative_labels = labels.chunk(2)
  148. # Generate labels
  149. base_loss = F.cross_entropy(
  150. token_logits.reshape(-1, token_logits.size(-1)),
  151. labels[:, 0].reshape(-1),
  152. ignore_index=-100,
  153. )
  154. codebook_labels = labels[:, 1 : 1 + self.model.config.num_codebooks].mT
  155. semantic_loss = F.cross_entropy(
  156. codebook_logits.reshape(-1, codebook_logits.size(-1)),
  157. codebook_labels.reshape(-1),
  158. ignore_index=-100,
  159. )
  160. loss = base_loss + semantic_loss
  161. # If we use dpo
  162. if self.use_dpo:
  163. negative_codebook_labels = negative_labels[
  164. :, 1 : 1 + self.model.config.num_codebooks
  165. ].mT
  166. positive_codebook_logps = self.get_batch_logps(
  167. codebook_logits, codebook_labels
  168. )
  169. negative_codebook_logps = self.get_batch_logps(
  170. negative_codebook_logits, negative_codebook_labels
  171. )
  172. # TODO: implement the reference model, avoid screwing up the gradients
  173. dpo_loss = -F.logsigmoid(
  174. (positive_codebook_logps - negative_codebook_logps) * self.dpo_beta
  175. ).mean()
  176. chosen_rewards = self.dpo_beta * positive_codebook_logps.detach()
  177. rejected_rewards = self.dpo_beta * negative_codebook_logps.detach()
  178. reward_accuracy = (chosen_rewards > rejected_rewards).float().mean()
  179. chosen_rewards, rejected_rewards = (
  180. chosen_rewards.mean(),
  181. rejected_rewards.mean(),
  182. )
  183. loss = loss + dpo_loss
  184. self.log(
  185. f"{stage}/dpo_loss",
  186. dpo_loss,
  187. on_step=True,
  188. on_epoch=False,
  189. prog_bar=False,
  190. logger=True,
  191. )
  192. self.log(
  193. f"{stage}/chosen_rewards",
  194. chosen_rewards,
  195. on_step=True,
  196. on_epoch=False,
  197. prog_bar=False,
  198. logger=True,
  199. )
  200. self.log(
  201. f"{stage}/rejected_rewards",
  202. rejected_rewards,
  203. on_step=True,
  204. on_epoch=False,
  205. prog_bar=False,
  206. logger=True,
  207. )
  208. self.log(
  209. f"{stage}/reward_accuracy",
  210. reward_accuracy,
  211. on_step=True,
  212. on_epoch=False,
  213. prog_bar=False,
  214. logger=True,
  215. )
  216. self.log(
  217. f"{stage}/loss",
  218. loss,
  219. on_step=True,
  220. on_epoch=False,
  221. prog_bar=True,
  222. logger=True,
  223. )
  224. self.log(
  225. f"{stage}/base_loss",
  226. base_loss,
  227. on_step=True,
  228. on_epoch=False,
  229. prog_bar=False,
  230. logger=True,
  231. )
  232. self.log(
  233. f"{stage}/semantic_loss",
  234. semantic_loss,
  235. on_step=True,
  236. on_epoch=False,
  237. prog_bar=False,
  238. logger=True,
  239. )
  240. # Top-5 accuracy
  241. accuracy = self.get_accuracy(codebook_logits, codebook_labels)
  242. self.log(
  243. f"{stage}/top_5_accuracy",
  244. accuracy,
  245. on_step=True,
  246. on_epoch=False,
  247. prog_bar=True,
  248. logger=True,
  249. )
  250. if self.model.config.num_codebooks != self.model.config.num_in_codebooks:
  251. accuracy = self.get_accuracy(
  252. codebook_logits[:, :, : self.model.config.num_in_codebooks],
  253. codebook_labels[:, :, : self.model.config.num_in_codebooks],
  254. )
  255. self.log(
  256. f"{stage}/top_5_accuracy_in",
  257. accuracy,
  258. on_step=True,
  259. on_epoch=False,
  260. prog_bar=True,
  261. logger=True,
  262. )
  263. return loss
  264. def get_accuracy(self, logits, labels):
  265. _, indices = logits.topk(5, dim=-1)
  266. correct = indices.eq(labels.unsqueeze(-1))
  267. correct[labels == -100] = 0
  268. correct = correct.sum()
  269. accuracy = correct / (labels != -100).sum()
  270. return accuracy
  271. def training_step(self, batch, batch_idx):
  272. return self._step(batch, batch_idx, "train")
  273. def validation_step(self, batch, batch_idx):
  274. return self._step(batch, batch_idx, "val")