spec_decode_base_sampler.py 8.8 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231
  1. from abc import abstractmethod
  2. from typing import Dict, Optional, Union
  3. import torch
  4. import torch.jit
  5. import torch.nn as nn
  6. class SpecDecodeBaseSampler(nn.Module):
  7. """Base class for samplers used for Speculative Decoding verification
  8. step.
  9. """
  10. def __init__(self, strict_mode: bool = False):
  11. """Base class constructor.
  12. Args:
  13. strict_mode: Whether or not to perform shape/device/dtype checks
  14. during sampling. This catches correctness issues but adds
  15. nontrivial latency.
  16. """
  17. super().__init__()
  18. self._strict_mode = strict_mode
  19. # NOTE: A "bonus token" is accepted iff all proposal tokens are
  20. # accepted. There is always only one possible bonus token. We store this
  21. # value in a variable for readability.
  22. self._num_bonus_tokens = 1
  23. self.num_accepted_tokens: Optional[torch.Tensor] = None
  24. self.num_emitted_tokens: Optional[torch.Tensor] = None
  25. self.num_draft_tokens: int = 0
  26. def init_gpu_tensors(self, device: Union[int, str]) -> None:
  27. assert self.num_accepted_tokens is None
  28. if isinstance(device, int):
  29. device = f"cuda:{device}"
  30. elif not isinstance(device, str):
  31. raise ValueError(f"Device must be int or str, get {type(device)}")
  32. self.num_accepted_tokens = torch.tensor(0,
  33. dtype=torch.long,
  34. device=device)
  35. self.num_emitted_tokens = torch.tensor(0,
  36. dtype=torch.long,
  37. device=device)
  38. @property
  39. def probs_dtype(self):
  40. return torch.float32
  41. @property
  42. def token_id_dtype(self):
  43. return torch.int64
  44. def _create_output(
  45. self,
  46. accepted: torch.Tensor, # [batch_size, k]
  47. substitute_token_ids: torch.Tensor, # [batch_size, k]
  48. draft_token_ids: torch.Tensor, # [batch_size, k]
  49. bonus_token_ids: torch.Tensor, # [batch_size]
  50. ) -> torch.Tensor:
  51. """Format output. Returns a matrix of token ids. When
  52. a token is rejected via sampling, all subsequent token ids are
  53. set to -1 for the sequence.
  54. Args:
  55. accepted: A boolean tensor indicating if the corresponding
  56. draft token in draft_token_ids should be accepted or not.
  57. substitute_token_ids: A tensor of token_ids that can be used
  58. as substitutes for the draft token ids if the proposed token
  59. is rejected.
  60. draft_token_ids: A tensor of token ids speculated by the
  61. draft model.
  62. bonus_token_ids: Token ids to use as the bonus token if
  63. all the draft tokens are accepted.
  64. Returns:
  65. A tensor containing the accepted token ids. The shape of the
  66. tensor is [batch_size, k + num_bonus_tokens]
  67. """
  68. batch_size, k = substitute_token_ids.shape
  69. bonus_token_ids = bonus_token_ids.squeeze()
  70. # Determine the index of the first False value for each row.
  71. limits = (accepted == 0).max(1).indices
  72. limits[~(accepted == 0).any(1)] = k
  73. # Create masks using the indices.
  74. indices = torch.arange(k, device=accepted.device).unsqueeze(0)
  75. accepted_mask = indices < limits.unsqueeze(1)
  76. after_false_mask = indices == limits.unsqueeze(1)
  77. # Create an extended output tensor
  78. output_with_bonus_tokens = -torch.ones(
  79. (batch_size, k + self._num_bonus_tokens),
  80. dtype=self.token_id_dtype,
  81. device=accepted.device)
  82. output = output_with_bonus_tokens[:, :k]
  83. # Fill in the first k columns of the output tensor using masks and data
  84. # tensors.
  85. output[:, :k] = torch.where(accepted_mask, draft_token_ids,
  86. -torch.ones_like(draft_token_ids))
  87. # Fill the last column.
  88. # We check output directly as accepted may have True values inconsistent
  89. # with causal acceptance.
  90. output_with_bonus_tokens[:, -1] = torch.where(output[:, -1] != -1,
  91. bonus_token_ids, -1)
  92. # Fill the recovered token ids.
  93. output.mul_(~after_false_mask).add_(
  94. substitute_token_ids.mul(after_false_mask))
  95. self.num_accepted_tokens += accepted.sum()
  96. self.num_emitted_tokens += (output_with_bonus_tokens != -1).sum()
  97. self.num_draft_tokens += batch_size * k
  98. return output_with_bonus_tokens
  99. def _raise_if_incorrect_input(
  100. self,
  101. target_probs: torch.Tensor,
  102. draft_token_ids: torch.Tensor,
  103. bonus_token_ids: torch.Tensor,
  104. draft_probs: Optional[torch.Tensor] = None,
  105. ) -> None:
  106. self._raise_if_incorrect_shape(target_probs, draft_token_ids,
  107. bonus_token_ids, draft_probs)
  108. self._raise_if_incorrect_dtype(target_probs, draft_token_ids,
  109. bonus_token_ids, draft_probs)
  110. self._raise_if_inconsistent_device(target_probs, draft_token_ids,
  111. bonus_token_ids, draft_probs)
  112. self._raise_if_out_of_bounds_vocab(target_probs.shape[-1],
  113. draft_token_ids, bonus_token_ids)
  114. def _raise_if_incorrect_shape(
  115. self,
  116. target_probs: torch.Tensor,
  117. draft_token_ids: torch.Tensor,
  118. bonus_token_ids: torch.Tensor,
  119. draft_probs: Optional[torch.Tensor] = None,
  120. ) -> None:
  121. (target_batch_size, num_target_probs,
  122. target_vocab_size) = target_probs.shape
  123. # validate the shape of draft token ids.
  124. draft_token_ids_batch_size, num_draft_token_ids = draft_token_ids.shape
  125. assert draft_token_ids_batch_size == target_batch_size
  126. assert num_draft_token_ids == num_target_probs
  127. # validate the shape of bonus token ids
  128. bonus_batch_size, num_bonus_tokens = bonus_token_ids.shape
  129. assert bonus_batch_size == target_batch_size
  130. assert num_bonus_tokens == self._num_bonus_tokens
  131. # validate the shape of draft probs if it is set
  132. if draft_probs is not None:
  133. (draft_batch_size, num_draft_probs,
  134. draft_vocab_size) = draft_probs.shape
  135. assert draft_batch_size == target_batch_size
  136. assert num_draft_probs == num_target_probs
  137. assert (draft_vocab_size == target_vocab_size
  138. ), f"{draft_vocab_size=} {target_vocab_size=}"
  139. def _raise_if_incorrect_dtype(
  140. self,
  141. target_probs: torch.Tensor,
  142. draft_token_ids: torch.Tensor,
  143. bonus_token_ids: torch.Tensor,
  144. draft_probs: Optional[torch.Tensor] = None,
  145. ) -> None:
  146. assert target_probs.dtype == self.probs_dtype
  147. assert draft_token_ids.dtype == self.token_id_dtype
  148. assert bonus_token_ids.dtype == self.token_id_dtype
  149. if draft_probs is not None:
  150. assert draft_probs.dtype == self.probs_dtype
  151. def _raise_if_inconsistent_device(
  152. self,
  153. target_probs: torch.Tensor,
  154. draft_token_ids: torch.Tensor,
  155. bonus_token_ids: torch.Tensor,
  156. draft_probs: Optional[torch.Tensor] = None,
  157. ) -> None:
  158. devices = [
  159. t.device for t in
  160. [target_probs, bonus_token_ids, draft_probs, draft_token_ids]
  161. if t is not None
  162. ]
  163. assert all([devices[0] == device for device in devices])
  164. def _raise_if_out_of_bounds_vocab(
  165. self,
  166. vocab_size: int,
  167. draft_token_ids: torch.Tensor,
  168. bonus_token_ids: torch.Tensor,
  169. ) -> None:
  170. assert torch.all(bonus_token_ids < vocab_size)
  171. assert torch.all(bonus_token_ids >= 0)
  172. assert torch.all(draft_token_ids < vocab_size)
  173. assert torch.all(draft_token_ids >= 0)
  174. class SpecDecodeDeterministicBaseSampler(SpecDecodeBaseSampler):
  175. """Base class for samplers used for Speculative Decoding verification
  176. step which are deterministic.
  177. """
  178. @abstractmethod
  179. def forward(
  180. self,
  181. target_probs: torch.Tensor,
  182. bonus_token_ids: torch.Tensor,
  183. draft_probs: torch.Tensor,
  184. draft_token_ids: torch.Tensor,
  185. ) -> torch.Tensor:
  186. raise NotImplementedError
  187. class SpecDecodeStochasticBaseSampler(SpecDecodeBaseSampler):
  188. """Base class for samplers used for Speculative Decoding verification
  189. step which are stochastic
  190. """
  191. @abstractmethod
  192. def forward(
  193. self,
  194. target_probs: torch.Tensor,
  195. bonus_token_ids: torch.Tensor,
  196. draft_probs: torch.Tensor,
  197. draft_token_ids: torch.Tensor,
  198. seeded_seqs: Optional[Dict[int, torch.Generator]] = None,
  199. ) -> torch.Tensor:
  200. raise NotImplementedError