block_manager_v2.py 9.6 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242
  1. """A block manager that manages token blocks."""
  2. from typing import Dict, List, Optional
  3. from aphrodite.processing.block.block_table import BlockTable
  4. from aphrodite.processing.block.cpu_gpu_block_allocator import (
  5. CpuGpuBlockAllocator)
  6. from aphrodite.processing.interfaces import AllocStatus, BlockSpaceManager
  7. from aphrodite.common.sequence import Sequence, SequenceGroup, SequenceStatus
  8. from aphrodite.common.utils import Device
  9. SeqId = int
  10. class BlockSpaceManagerV2(BlockSpaceManager):
  11. """BlockSpaceManager which manages the allocation of KV cache.
  12. It owns responsibility for allocation, swapping, allocating memory for
  13. autoregressively-generated tokens, and other advanced features such as
  14. prefix caching, forking/copy-on-write, and sliding-window memory allocation.
  15. The current implementation is partial; in particular prefix caching and
  16. sliding-window are not feature complete.
  17. Lookahead slots
  18. The block manager has the notion of a "lookahead slot". These are slots
  19. in the KV cache that are allocated for a sequence. Unlike the other
  20. allocated slots, the content of these slots is undefined -- the worker
  21. may use the memory allocations in any way.
  22. In practice, a worker could use these lookahead slots to run multiple
  23. forward passes for a single scheduler invocation. Each successive
  24. forward pass would write KV activations to the corresponding lookahead
  25. slot. This allows low inter-token latency use-cases, where the overhead
  26. of continuous batching scheduling is amortized over >1 generated tokens.
  27. Speculative decoding uses lookahead slots to store KV activations of
  28. proposal tokens.
  29. Args:
  30. block_size (int): The size of each memory block.
  31. num_gpu_blocks (int): The number of memory blocks allocated on GPU.
  32. num_cpu_blocks (int): The number of memory blocks allocated on CPU.
  33. watermark (float, optional): The threshold used for memory swapping.
  34. Defaults to 0.01.
  35. sliding_window (Optional[int], optional): The size of the sliding
  36. window. Defaults to None.
  37. enable_caching (bool, optional): Flag indicating whether caching is
  38. enabled. Defaults to False.
  39. """
  40. def __init__(
  41. self,
  42. block_size: int,
  43. num_gpu_blocks: int,
  44. num_cpu_blocks: int,
  45. watermark: float = 0.01,
  46. sliding_window: Optional[int] = None,
  47. enable_caching: bool = False,
  48. ) -> None:
  49. self.block_size = block_size
  50. self.num_total_gpu_blocks = num_gpu_blocks
  51. self.num_total_cpu_blocks = num_cpu_blocks
  52. assert sliding_window is None, "Sliding window not yet supported"
  53. self.block_sliding_window = None
  54. self.watermark = watermark
  55. assert watermark >= 0.0
  56. assert not enable_caching, "Prefix caching not yet supported"
  57. self.enable_caching = enable_caching
  58. self.watermark_blocks = int(watermark * num_gpu_blocks)
  59. self.block_allocator = CpuGpuBlockAllocator.create(
  60. # Currently, only naive blocks are supported (no prefix caching).
  61. allocator_type="naive",
  62. num_gpu_blocks=num_gpu_blocks,
  63. num_cpu_blocks=num_cpu_blocks,
  64. block_size=block_size,
  65. )
  66. self.block_tables: Dict[SeqId, BlockTable] = {}
  67. def can_allocate(self, seq_group: SequenceGroup) -> AllocStatus:
  68. # FIXME: Here we assume that all sequences in the group share
  69. # the same prompt. This may not be true for preempted sequences.
  70. seq = seq_group.get_seqs(status=SequenceStatus.WAITING)[0]
  71. num_required_blocks = BlockTable.get_num_required_blocks(
  72. seq.get_token_ids(),
  73. block_size=self.block_size,
  74. )
  75. assert self.block_sliding_window is None
  76. if self.block_sliding_window is not None:
  77. num_required_blocks = min(num_required_blocks,
  78. self.block_sliding_window)
  79. num_free_gpu_blocks = self.block_allocator.get_num_free_blocks(
  80. device=Device.GPU)
  81. # Use watermark to avoid frequent cache eviction.
  82. if (self.num_total_gpu_blocks - num_required_blocks <
  83. self.watermark_blocks):
  84. return AllocStatus.NEVER
  85. if num_free_gpu_blocks - num_required_blocks >= self.watermark_blocks:
  86. return AllocStatus.OK
  87. else:
  88. return AllocStatus.LATER
  89. def allocate(self, seq_group: SequenceGroup) -> None:
  90. waiting_seqs = seq_group.get_seqs(status=SequenceStatus.WAITING)
  91. assert not (set(seq.seq_id for seq in waiting_seqs)
  92. & self.block_tables.keys()), "block table already exists"
  93. # NOTE: Here we assume that all sequences in the group have the same
  94. # prompt.
  95. seq = waiting_seqs[0]
  96. block_table = BlockTable(
  97. block_size=self.block_size,
  98. block_allocator=self.block_allocator,
  99. )
  100. assert self.block_sliding_window is None
  101. block_table.allocate(seq.get_token_ids())
  102. self.block_tables[seq.seq_id] = block_table
  103. # Assign the block table for each sequence.
  104. for seq in waiting_seqs[1:]:
  105. self.block_tables[seq.seq_id] = block_table.fork()
  106. def can_append_slots(self, seq_group: SequenceGroup,
  107. num_lookahead_slots: int) -> bool:
  108. """Determine if there is enough space in the GPU KV cache to continue
  109. generation of the specified sequence group.
  110. We use a worst-case heuristic: assume each touched block will require a
  111. new allocation (either via CoW or new block). We can append slots if the
  112. number of touched blocks is less than the number of free blocks.
  113. "Lookahead slots" are slots that are allocated in addition to the slots
  114. for known tokens. The contents of the lookahead slots are not defined.
  115. This is used by speculative decoding when speculating future tokens.
  116. """
  117. num_touched_blocks = 0
  118. for seq in seq_group.get_seqs(status=SequenceStatus.RUNNING):
  119. block_table = self.block_tables[seq.seq_id]
  120. num_touched_blocks += (
  121. block_table.get_num_blocks_touched_by_append_slots(
  122. token_ids=block_table.get_unseen_token_ids(
  123. seq.get_token_ids()),
  124. num_lookahead_slots=num_lookahead_slots,
  125. ))
  126. num_free_gpu_blocks = self.block_allocator.get_num_free_blocks(
  127. Device.GPU)
  128. return num_touched_blocks <= num_free_gpu_blocks
  129. def append_slots(
  130. self,
  131. seq: Sequence,
  132. num_lookahead_slots: int,
  133. ) -> Dict[int, List[int]]:
  134. block_table = self.block_tables[seq.seq_id]
  135. block_table.append_token_ids(
  136. token_ids=block_table.get_unseen_token_ids(seq.get_token_ids()),
  137. num_lookahead_slots=num_lookahead_slots,
  138. )
  139. # Return any new copy-on-writes.
  140. new_cows = self.block_allocator.clear_copy_on_writes()
  141. return new_cows
  142. def free(self, seq: Sequence) -> None:
  143. if seq.seq_id not in self.block_tables:
  144. # Already freed or haven't been scheduled yet.
  145. return
  146. self.block_tables[seq.seq_id].free()
  147. del self.block_tables[seq.seq_id]
  148. def get_block_table(self, seq: Sequence) -> List[int]:
  149. assert seq.seq_id in self.block_tables
  150. block_ids = self.block_tables[seq.seq_id].physical_block_ids
  151. assert all(b is not None for b in block_ids)
  152. return block_ids
  153. def access_all_blocks_in_seq(self, seq, now):
  154. # TODO add prefix caching support.
  155. pass
  156. def mark_blocks_as_computed(self, seq_group: SequenceGroup):
  157. # We ignore the sequence group as its not necessary. After the batch is
  158. # formed by the scheduler, we do not need to mark blocks from individual
  159. # sequence groups as computed -- all blocks in the batch can be marked
  160. # as computed.
  161. self.block_allocator.mark_blocks_as_computed()
  162. def get_common_computed_block_ids(self, seqs: List[Sequence]) -> List[int]:
  163. """Determine which blocks for which we skip prefill.
  164. With prefix caching we can skip prefill for previously-generated blocks.
  165. Currently, the attention implementation only supports skipping cached
  166. blocks if they are a contiguous prefix of cached blocks.
  167. This method determines which blocks can be safely skipped for all
  168. sequences in the sequence group.
  169. """
  170. seq_block_ids = [
  171. self.block_tables[seq.seq_id].physical_block_ids for seq in seqs
  172. ]
  173. return self.block_allocator.get_common_computed_block_ids(
  174. seq_block_ids)
  175. def fork(self, parent_seq: Sequence, child_seq: Sequence) -> None:
  176. src_block_table = self.block_tables[parent_seq.seq_id]
  177. self.block_tables[child_seq.seq_id] = src_block_table.fork()
  178. def can_swap_in(self, seq_group: SequenceGroup,
  179. num_lookahead_slots: int) -> bool:
  180. return False
  181. def swap_in(self, seq_group: SequenceGroup,
  182. num_lookahead_slots: int) -> Dict[int, int]:
  183. raise NotImplementedError
  184. def can_swap_out(self, seq_group: SequenceGroup) -> bool:
  185. return False
  186. def swap_out(self, seq_group: SequenceGroup) -> Dict[int, int]:
  187. raise NotImplementedError
  188. def get_num_free_gpu_blocks(self) -> int:
  189. return self.block_allocator.get_num_free_blocks(Device.GPU)
  190. def get_num_free_cpu_blocks(self) -> int:
  191. return self.block_allocator.get_num_free_blocks(Device.CPU)