scheduler.py 61 KB

12345678910111213141516171819202122232425262728293031323334353637383940414243444546474849505152535455565758596061626364656667686970717273747576777879808182838485868788899091929394959697989910010110210310410510610710810911011111211311411511611711811912012112212312412512612712812913013113213313413513613713813914014114214314414514614714814915015115215315415515615715815916016116216316416516616716816917017117217317417517617717817918018118218318418518618718818919019119219319419519619719819920020120220320420520620720820921021121221321421521621721821922022122222322422522622722822923023123223323423523623723823924024124224324424524624724824925025125225325425525625725825926026126226326426526626726826927027127227327427527627727827928028128228328428528628728828929029129229329429529629729829930030130230330430530630730830931031131231331431531631731831932032132232332432532632732832933033133233333433533633733833934034134234334434534634734834935035135235335435535635735835936036136236336436536636736836937037137237337437537637737837938038138238338438538638738838939039139239339439539639739839940040140240340440540640740840941041141241341441541641741841942042142242342442542642742842943043143243343443543643743843944044144244344444544644744844945045145245345445545645745845946046146246346446546646746846947047147247347447547647747847948048148248348448548648748848949049149249349449549649749849950050150250350450550650750850951051151251351451551651751851952052152252352452552652752852953053153253353453553653753853954054154254354454554654754854955055155255355455555655755855956056156256356456556656756856957057157257357457557657757857958058158258358458558658758858959059159259359459559659759859960060160260360460560660760860961061161261361461561661761861962062162262362462562662762862963063163263363463563663763863964064164264364464564664764864965065165265365465565665765865966066166266366466566666766866967067167267367467567667767867968068168268368468568668768868969069169269369469569669769869970070170270370470570670770870971071171271371471571671771871972072172272372472572672772872973073173273373473573673773873974074174274374474574674774874975075175275375475575675775875976076176276376476576676776876977077177277377477577677777877978078178278378478578678778878979079179279379479579679779879980080180280380480580680780880981081181281381481581681781881982082182282382482582682782882983083183283383483583683783883984084184284384484584684784884985085185285385485585685785885986086186286386486586686786886987087187287387487587687787887988088188288388488588688788888989089189289389489589689789889990090190290390490590690790890991091191291391491591691791891992092192292392492592692792892993093193293393493593693793893994094194294394494594694794894995095195295395495595695795895996096196296396496596696796896997097197297397497597697797897998098198298398498598698798898999099199299399499599699799899910001001100210031004100510061007100810091010101110121013101410151016101710181019102010211022102310241025102610271028102910301031103210331034103510361037103810391040104110421043104410451046104710481049105010511052105310541055105610571058105910601061106210631064106510661067106810691070107110721073107410751076107710781079108010811082108310841085108610871088108910901091109210931094109510961097109810991100110111021103110411051106110711081109111011111112111311141115111611171118111911201121112211231124112511261127112811291130113111321133113411351136113711381139114011411142114311441145114611471148114911501151115211531154115511561157115811591160116111621163116411651166116711681169117011711172117311741175117611771178117911801181118211831184118511861187118811891190119111921193119411951196119711981199120012011202120312041205120612071208120912101211121212131214121512161217121812191220122112221223122412251226122712281229123012311232123312341235123612371238123912401241124212431244124512461247124812491250125112521253125412551256125712581259126012611262126312641265126612671268126912701271127212731274127512761277127812791280128112821283128412851286128712881289129012911292129312941295129612971298129913001301130213031304130513061307130813091310131113121313131413151316131713181319132013211322132313241325132613271328132913301331133213331334133513361337133813391340134113421343134413451346134713481349135013511352135313541355135613571358135913601361136213631364136513661367136813691370137113721373137413751376137713781379138013811382138313841385138613871388138913901391139213931394139513961397139813991400140114021403140414051406140714081409141014111412141314141415141614171418141914201421142214231424142514261427142814291430
  1. import enum
  2. import os
  3. import random
  4. import time
  5. from collections import deque
  6. from dataclasses import dataclass, field
  7. from typing import Deque, Dict, Iterable, List, Optional, Set, Tuple, Union
  8. from loguru import logger
  9. from aphrodite.common.config import CacheConfig, LoRAConfig, SchedulerConfig
  10. from aphrodite.common.sequence import (Sequence, SequenceData, SequenceGroup,
  11. SequenceGroupMetadata,
  12. SequenceGroupMetadataDelta,
  13. SequenceStatus)
  14. from aphrodite.common.utils import Device, PyObjectCache
  15. from aphrodite.lora.request import LoRARequest
  16. from aphrodite.processing.interfaces import AllocStatus, BlockSpaceManager
  17. from aphrodite.prompt_adapter.request import PromptAdapterRequest
  18. # Test-only. If configured, decode is preempted with
  19. # ARTIFICIAL_PREEMPTION_PROB% probability.
  20. ENABLE_ARTIFICIAL_PREEMPT = bool(
  21. os.getenv("APHRODITE_TEST_ENABLE_ARTIFICIAL_PREEMPT", False)) # noqa
  22. ARTIFICIAL_PREEMPTION_PROB = 0.5
  23. ARTIFICIAL_PREEMPTION_MAX_CNT = 500
  24. class PreemptionMode(enum.Enum):
  25. """Preemption modes.
  26. 1. Swapping: Swap out the blocks of the preempted sequences to CPU memory
  27. and swap them back in when the sequences are resumed.
  28. 2. Recomputation: Discard the blocks of the preempted sequences and
  29. recompute them when the sequences are resumed, treating the sequences as
  30. new prompts.
  31. """
  32. SWAP = enum.auto()
  33. RECOMPUTE = enum.auto()
  34. @dataclass
  35. class SchedulingBudget:
  36. """The available slots for scheduling.
  37. TODO: Right now, the budget is request_id-aware meaning it can ignore
  38. budget update from the same request_id. It is because in normal scheduling
  39. path, we update RUNNING num_seqs ahead of time, meaning it could be
  40. updated more than once when scheduling RUNNING requests. Since this won't
  41. happen if we only have chunked prefill scheduling, we can remove this
  42. feature from the API when chunked prefill is enabled by default.
  43. """
  44. token_budget: int
  45. max_num_seqs: int
  46. _request_ids_num_batched_tokens: Set[str] = field(default_factory=set)
  47. _request_ids_num_curr_seqs: Set[str] = field(default_factory=set)
  48. _num_batched_tokens: int = 0
  49. _num_curr_seqs: int = 0
  50. def can_schedule(self, *, num_new_tokens: int, num_new_seqs: int):
  51. assert num_new_tokens != 0
  52. assert num_new_seqs != 0
  53. return (self.num_batched_tokens + num_new_tokens <= self.token_budget
  54. and self.num_curr_seqs + num_new_seqs <= self.max_num_seqs)
  55. def remaining_token_budget(self):
  56. return self.token_budget - self.num_batched_tokens
  57. def add_num_batched_tokens(self, req_id: str, num_batched_tokens: int):
  58. if req_id in self._request_ids_num_batched_tokens:
  59. return
  60. self._request_ids_num_batched_tokens.add(req_id)
  61. self._num_batched_tokens += num_batched_tokens
  62. def subtract_num_batched_tokens(self, req_id: str,
  63. num_batched_tokens: int):
  64. if req_id in self._request_ids_num_batched_tokens:
  65. self._request_ids_num_batched_tokens.remove(req_id)
  66. self._num_batched_tokens -= num_batched_tokens
  67. def add_num_seqs(self, req_id: str, num_curr_seqs: int):
  68. if req_id in self._request_ids_num_curr_seqs:
  69. return
  70. self._request_ids_num_curr_seqs.add(req_id)
  71. self._num_curr_seqs += num_curr_seqs
  72. def subtract_num_seqs(self, req_id: str, num_curr_seqs: int):
  73. if req_id in self._request_ids_num_curr_seqs:
  74. self._request_ids_num_curr_seqs.remove(req_id)
  75. self._num_curr_seqs -= num_curr_seqs
  76. @property
  77. def num_batched_tokens(self):
  78. return self._num_batched_tokens
  79. @property
  80. def num_curr_seqs(self):
  81. return self._num_curr_seqs
  82. @dataclass
  83. class ScheduledSequenceGroup:
  84. # A sequence group that's scheduled.
  85. seq_group: SequenceGroup
  86. # The total chunk size (number of tokens) to process for next iteration.
  87. # 1 for decoding. Same as prompt tokens for prefill, but if prefill is
  88. # chunked, it can be smaller than that.
  89. token_chunk_size: int
  90. negative_token_chunk_size: int = 0
  91. @dataclass
  92. class SchedulerOutputs:
  93. """The scheduling decision made from a scheduler."""
  94. # Scheduled sequence groups.
  95. scheduled_seq_groups: Iterable[ScheduledSequenceGroup]
  96. # Number of prefill groups scheduled.
  97. num_prefill_groups: int
  98. # Total number of batched tokens.
  99. num_batched_tokens: int
  100. # Blocks to swap in. List of CPU -> GPU block number.
  101. blocks_to_swap_in: List[Tuple[int, int]]
  102. # Blocks to swap out. List of GPU -> CPU block number.
  103. blocks_to_swap_out: List[Tuple[int, int]]
  104. # Blocks to copy. Source to dest block.
  105. blocks_to_copy: List[Tuple[int, int]]
  106. # Sequence groups that are going to be ignored.
  107. ignored_seq_groups: List[SequenceGroup]
  108. # The number of slots for lookahead decoding.
  109. num_lookahead_slots: int
  110. # The number of requests in the running queue
  111. running_queue_size: int
  112. preempted: int
  113. def __post_init__(self):
  114. # Swap in and swap out should never happen at the same time.
  115. assert not (self.blocks_to_swap_in and self.blocks_to_swap_out)
  116. self.num_loras: int = len(self.lora_requests)
  117. if self.num_loras > 0:
  118. self._sort_by_lora_ids()
  119. self.num_prompt_adapters: int = len(self.prompt_adapter_requests)
  120. def is_empty(self) -> bool:
  121. # NOTE: We do not consider the ignored sequence groups.
  122. return (not self.scheduled_seq_groups and not self.blocks_to_swap_in
  123. and not self.blocks_to_swap_out and not self.blocks_to_copy)
  124. def _sort_by_lora_ids(self):
  125. assert 0 <= self.num_prefill_groups <= len(self.scheduled_seq_groups)
  126. def key_fn(group: ScheduledSequenceGroup):
  127. key = (group.seq_group.lora_int_id, group.seq_group.request_id)
  128. if 0 < self.num_prefill_groups < len(self.scheduled_seq_groups):
  129. # Sort sequence groups so that all prefills come before all
  130. # decodes as required by chunked prefill.
  131. return (not group.seq_group.is_prefill(), *key)
  132. return key
  133. self.scheduled_seq_groups = sorted(self.scheduled_seq_groups,
  134. key=key_fn)
  135. @property
  136. def lora_requests(self) -> Set[LoRARequest]:
  137. return {
  138. g.seq_group.lora_request
  139. for g in self.scheduled_seq_groups
  140. if g.seq_group.lora_request is not None
  141. }
  142. @property
  143. def prompt_adapter_requests(self) -> Set[PromptAdapterRequest]:
  144. return {
  145. g.seq_group.prompt_adapter_request
  146. for g in self.scheduled_seq_groups
  147. if g.seq_group.prompt_adapter_request is not None
  148. }
  149. @dataclass
  150. class SchedulerRunningOutputs:
  151. """The requests that are scheduled from a running queue.
  152. Could contain prefill (prefill that's chunked) or decodes. If there's not
  153. enough memory, it can be preempted (for recompute) or swapped out.
  154. """
  155. # Selected sequences that are running and in a decoding phase.
  156. decode_seq_groups: List[ScheduledSequenceGroup]
  157. # Selected sequences that are running and in a prefill phase.
  158. # I.e., it means the prefill has been chunked.
  159. prefill_seq_groups: List[ScheduledSequenceGroup]
  160. # The preempted sequences.
  161. preempted: List[SequenceGroup]
  162. # Sequences that are swapped out.
  163. swapped_out: List[SequenceGroup]
  164. # The blocks to swap out.
  165. blocks_to_swap_out: List[Tuple[int, int]]
  166. # The blocks to copy.
  167. blocks_to_copy: List[Tuple[int, int]]
  168. # The number of slots for lookahead decoding.
  169. num_lookahead_slots: int
  170. # Optimization for fast-access to seq_group lists
  171. decode_seq_groups_list: List[SequenceGroup]
  172. prefill_seq_groups_list: List[SequenceGroup]
  173. @classmethod
  174. def create_empty(cls) -> "SchedulerRunningOutputs":
  175. return SchedulerRunningOutputs(
  176. decode_seq_groups=[],
  177. prefill_seq_groups=[],
  178. preempted=[],
  179. swapped_out=[],
  180. blocks_to_swap_out=[],
  181. blocks_to_copy=[],
  182. num_lookahead_slots=0,
  183. decode_seq_groups_list=[],
  184. prefill_seq_groups_list=[],
  185. )
  186. @dataclass
  187. class SchedulerSwappedInOutputs:
  188. """The requests that are scheduled from a swap queue.
  189. Could contain prefill (prefill that's chunked) or decodes.
  190. """
  191. # Selected sequences that are going to be swapped in and is in a
  192. # decoding phase.
  193. decode_seq_groups: List[SequenceGroup]
  194. # Selected sequences that are going to be swapped in and in a prefill
  195. # phase. I.e., it means the prefill has been chunked.
  196. prefill_seq_groups: List[SequenceGroup]
  197. # The blocks to swap in.
  198. blocks_to_swap_in: List[Tuple[int, int]]
  199. # The blocks to copy.
  200. blocks_to_copy: List[Tuple[int, int]]
  201. # The number of slots for lookahead decoding.
  202. num_lookahead_slots: int
  203. # Infeasible sequence groups.
  204. infeasible_seq_groups: List[SequenceGroup]
  205. @classmethod
  206. def create_empty(cls) -> "SchedulerSwappedInOutputs":
  207. return SchedulerSwappedInOutputs(
  208. decode_seq_groups=[],
  209. prefill_seq_groups=[],
  210. blocks_to_swap_in=[],
  211. blocks_to_copy=[],
  212. num_lookahead_slots=0,
  213. infeasible_seq_groups=[],
  214. )
  215. @dataclass
  216. class SchedulerPrefillOutputs:
  217. """The requests that are scheduled from a waiting queue.
  218. Could contain a fresh prefill requests or preempted requests that need
  219. to be recomputed from scratch.
  220. """
  221. # Selected sequences for prefill.
  222. seq_groups: List[SequenceGroup]
  223. # Ignored sequence groups.
  224. ignored_seq_groups: List[SequenceGroup]
  225. num_lookahead_slots: int
  226. @classmethod
  227. def create_empty(cls) -> "SchedulerPrefillOutputs":
  228. return SchedulerPrefillOutputs(
  229. seq_groups=[],
  230. ignored_seq_groups=[],
  231. num_lookahead_slots=0,
  232. )
  233. def seq_group_metadata_builder():
  234. return SequenceGroupMetadata(request_id="",
  235. is_prompt=False,
  236. seq_data={},
  237. sampling_params=None,
  238. block_tables={})
  239. def scheduler_running_outputs_builder():
  240. return SchedulerRunningOutputs(decode_seq_groups=[],
  241. prefill_seq_groups=[],
  242. preempted=[],
  243. swapped_out=[],
  244. blocks_to_swap_out=[],
  245. blocks_to_copy=[],
  246. num_lookahead_slots=0,
  247. prefill_seq_groups_list=[],
  248. decode_seq_groups_list=[])
  249. def scheduled_seq_group_builder():
  250. return ScheduledSequenceGroup(
  251. seq_group=None, token_chunk_size=0, negative_token_chunk_size=0)
  252. class Scheduler:
  253. def __init__(
  254. self,
  255. scheduler_config: SchedulerConfig,
  256. cache_config: CacheConfig,
  257. lora_config: Optional[LoRAConfig],
  258. pipeline_parallel_size: int = 1,
  259. ) -> None:
  260. self.scheduler_config = scheduler_config
  261. self.cache_config = cache_config
  262. # Note for LoRA scheduling: the current policy is extremely
  263. # simple and NOT fair. It can lead to starvation of some
  264. # LoRAs. This should be improved in the future.
  265. self.lora_config = lora_config
  266. version = "v1"
  267. if self.scheduler_config.use_v2_block_manager:
  268. version = "v2"
  269. if (self.scheduler_config.embedding_mode
  270. or self.scheduler_config.is_attention_free):
  271. version = "placeholder"
  272. BlockSpaceManagerImpl = BlockSpaceManager.get_block_space_manager_class(
  273. version)
  274. num_gpu_blocks = cache_config.num_gpu_blocks
  275. if num_gpu_blocks:
  276. num_gpu_blocks //= pipeline_parallel_size
  277. num_cpu_blocks = cache_config.num_cpu_blocks
  278. if num_cpu_blocks:
  279. num_cpu_blocks //= pipeline_parallel_size
  280. # Create the block space manager.
  281. self.block_manager = BlockSpaceManagerImpl(
  282. block_size=self.cache_config.block_size,
  283. num_gpu_blocks=num_gpu_blocks,
  284. num_cpu_blocks=num_cpu_blocks,
  285. sliding_window=self.cache_config.sliding_window,
  286. enable_caching=self.cache_config.enable_prefix_caching)
  287. # Sequence groups in the WAITING state.
  288. # Contain new prefill or preempted requests.
  289. self.waiting: Deque[SequenceGroup] = deque()
  290. # Sequence groups in the RUNNING state.
  291. # Contain decode requests.
  292. self.running: Deque[SequenceGroup] = deque()
  293. # Sequence groups in the SWAPPED state.
  294. # Contain decode requests that are swapped out.
  295. self.swapped: Deque[SequenceGroup] = deque()
  296. # Sequence groups finished requests ids since last step iteration.
  297. # It lets the model know that any state associated with these requests
  298. # can and must be released after the current step.
  299. # This is used to evict the finished requests from the Mamba cache.
  300. self._finished_requests_ids: List[str] = list()
  301. # Time at previous scheduling step
  302. self.prev_time = 0.0
  303. # Did we schedule a prompt at previous step?
  304. self.prev_prompt = False
  305. # Latency of the last prompt step
  306. self.last_prompt_latency = 0.0
  307. # preemption mode, RECOMPUTE or SWAP
  308. self.user_specified_preemption_mode = scheduler_config.preemption_mode
  309. # The following field is test-only. It is used to inject artificial
  310. # preemption.
  311. self.enable_artificial_preemption = ENABLE_ARTIFICIAL_PREEMPT
  312. self.artificial_preempt_cnt = (ARTIFICIAL_PREEMPTION_MAX_CNT
  313. if self.enable_artificial_preemption
  314. else 0)
  315. self.num_cumulative_preemption: int = 0
  316. # Used to cache python objects
  317. self._scheduler_running_outputs_cache: PyObjectCache = PyObjectCache(
  318. scheduler_running_outputs_builder)
  319. self._scheduled_seq_group_cache: PyObjectCache = PyObjectCache(
  320. scheduled_seq_group_builder)
  321. @property
  322. def lora_enabled(self) -> bool:
  323. return bool(self.lora_config)
  324. @property
  325. def num_decoding_tokens_per_seq(self) -> int:
  326. """The number of new tokens."""
  327. return 1
  328. def add_seq_group(self, seq_group: SequenceGroup) -> None:
  329. # Add sequence groups to the waiting queue.
  330. self.waiting.append(seq_group)
  331. def _add_seq_group_to_running(self, seq_group: SequenceGroup) -> None:
  332. # Add sequence groups to the running queue.
  333. # Only for testing purposes.
  334. self.running.append(seq_group)
  335. def _add_seq_group_to_swapped(self, seq_group: SequenceGroup) -> None:
  336. # Add sequence groups to the swapped queue.
  337. # Only for testing purposes.
  338. self.swapped.append(seq_group)
  339. def abort_seq_group(self, request_id: Union[str, Iterable[str]]) -> None:
  340. """Aborts a sequence group with the given ID.
  341. Check if the sequence group with the given ID
  342. is present in any of the state queue.
  343. If present, remove the sequence group from the state queue.
  344. Also, if any of the sequences in the sequence group is not finished,
  345. free the sequence with status `FINISHED_ABORTED`.
  346. Otherwise, do nothing.
  347. Args:
  348. request_id: The ID(s) of the sequence group to abort.
  349. """
  350. if isinstance(request_id, str):
  351. request_id = (request_id, )
  352. request_ids = set(request_id)
  353. for state_queue in [self.waiting, self.running, self.swapped]:
  354. aborted_groups: List[SequenceGroup] = []
  355. for seq_group in state_queue:
  356. if not request_ids:
  357. # Using 'break' here may add two extra iterations,
  358. # but is acceptable to reduce complexity.
  359. break
  360. if seq_group.request_id in request_ids:
  361. # Appending aborted group into pending list.
  362. aborted_groups.append(seq_group)
  363. request_ids.remove(seq_group.request_id)
  364. for aborted_group in aborted_groups:
  365. # Remove the sequence group from the state queue.
  366. state_queue.remove(aborted_group)
  367. # Remove the aborted request from the Mamba cache.
  368. self._finished_requests_ids.append(aborted_group.request_id)
  369. for seq in aborted_group.get_seqs():
  370. if seq.is_finished():
  371. continue
  372. seq.status = SequenceStatus.FINISHED_ABORTED
  373. self.free_seq(seq)
  374. self._free_seq_group_cross_attn_blocks(aborted_group)
  375. def _free_seq_group_cross_attn_blocks(
  376. self,
  377. seq_group: SequenceGroup,
  378. ) -> None:
  379. """
  380. Free a sequence group from a cross-attention block table.
  381. Has no effect on decoder-only models.
  382. """
  383. if seq_group.is_encoder_decoder():
  384. self.block_manager.free_cross(seq_group)
  385. def has_unfinished_seqs(self) -> bool:
  386. return len(self.waiting) != 0 or len(self.running) != 0 or len(
  387. self.swapped) != 0
  388. def get_prefix_cache_hit_rate(self, device: Device) -> float:
  389. return self.block_manager.get_prefix_cache_hit_rate(device)
  390. def get_num_unfinished_seq_groups(self) -> int:
  391. return len(self.waiting) + len(self.running) + len(self.swapped)
  392. def get_and_reset_finished_requests_ids(self) -> List[str]:
  393. """Flushes the list of request ids of previously finished seq_groups."""
  394. finished_requests_ids = self._finished_requests_ids
  395. self._finished_requests_ids = list()
  396. return finished_requests_ids
  397. def _schedule_running(
  398. self,
  399. budget: SchedulingBudget,
  400. curr_loras: Optional[Set[int]],
  401. enable_chunking: bool = False,
  402. ) -> SchedulerRunningOutputs:
  403. """Schedule sequence groups that are running.
  404. Running queue should include decode and chunked prefill requests.
  405. Args:
  406. budget: The scheduling budget. The argument is in-place updated
  407. when any decodes are preempted.
  408. curr_loras: Currently batched lora request ids. The argument is
  409. in-place updated when any decodes are preempted.
  410. enable_chunking: If True, seq group can be chunked and only a
  411. chunked number of tokens are scheduled if
  412. `budget.num_batched_tokens` has not enough capacity to schedule
  413. all tokens.
  414. Returns:
  415. SchedulerRunningOutputs.
  416. """
  417. ret: SchedulerRunningOutputs = \
  418. self._scheduler_running_outputs_cache.get_object()
  419. ret.blocks_to_swap_out.clear()
  420. ret.blocks_to_copy.clear()
  421. ret.decode_seq_groups.clear()
  422. ret.prefill_seq_groups.clear()
  423. ret.preempted.clear()
  424. ret.swapped_out.clear()
  425. ret.num_lookahead_slots = self._get_num_lookahead_slots(
  426. is_prefill=False)
  427. ret.decode_seq_groups_list.clear()
  428. ret.prefill_seq_groups_list.clear()
  429. # Blocks that need to be swapped or copied before model execution.
  430. blocks_to_swap_out: List[Tuple[int, int]] = ret.blocks_to_swap_out
  431. blocks_to_copy: List[Tuple[int, int]] = ret.blocks_to_copy
  432. decode_seq_groups: List[ScheduledSequenceGroup] = ret.decode_seq_groups
  433. prefill_seq_groups: List[
  434. ScheduledSequenceGroup] = ret.prefill_seq_groups
  435. preempted: List[SequenceGroup] = ret.preempted
  436. swapped_out: List[SequenceGroup] = ret.swapped_out
  437. # NOTE: Preemption happens only when there is no available slot
  438. # to keep all the sequence groups in the RUNNING state.
  439. running_queue = self.running
  440. while running_queue:
  441. seq_group = running_queue[0]
  442. num_running_tokens, negative_num_running_tokens = (
  443. self._get_num_new_tokens(
  444. seq_group, SequenceStatus.RUNNING,
  445. enable_chunking, budget))
  446. if num_running_tokens + negative_num_running_tokens == 0:
  447. break
  448. running_queue.popleft()
  449. while not self._can_append_slots(seq_group):
  450. budget.subtract_num_batched_tokens(seq_group.request_id,
  451. num_running_tokens +
  452. negative_num_running_tokens)
  453. num_running_seqs = seq_group.get_max_num_running_seqs()
  454. budget.subtract_num_seqs(seq_group.request_id,
  455. num_running_seqs)
  456. if (curr_loras is not None and seq_group.lora_int_id > 0
  457. and seq_group.lora_int_id in curr_loras):
  458. curr_loras.remove(seq_group.lora_int_id)
  459. if running_queue:
  460. # Preempt the lowest-priority sequence groups.
  461. victim_seq_group = running_queue.pop()
  462. preempted_mode = self._preempt(victim_seq_group,
  463. blocks_to_swap_out)
  464. if preempted_mode == PreemptionMode.RECOMPUTE:
  465. preempted.append(victim_seq_group)
  466. else:
  467. swapped_out.append(victim_seq_group)
  468. else:
  469. # No other sequence groups can be preempted.
  470. # Preempt the current sequence group.
  471. preempted_mode = self._preempt(seq_group,
  472. blocks_to_swap_out)
  473. if preempted_mode == PreemptionMode.RECOMPUTE:
  474. preempted.append(seq_group)
  475. else:
  476. swapped_out.append(seq_group)
  477. break
  478. else:
  479. self._append_slots(seq_group, blocks_to_copy)
  480. is_prefill = seq_group.is_prefill()
  481. scheduled_seq_group: ScheduledSequenceGroup = \
  482. self._scheduled_seq_group_cache.get_object()
  483. scheduled_seq_group.seq_group = seq_group
  484. if is_prefill:
  485. scheduled_seq_group.token_chunk_size = num_running_tokens
  486. scheduled_seq_group.negative_token_chunk_size = \
  487. negative_num_running_tokens
  488. prefill_seq_groups.append(scheduled_seq_group)
  489. ret.prefill_seq_groups_list.append(seq_group)
  490. else:
  491. scheduled_seq_group.token_chunk_size = 1
  492. scheduled_seq_group.negative_token_chunk_size = 1
  493. decode_seq_groups.append(scheduled_seq_group)
  494. ret.decode_seq_groups_list.append(seq_group)
  495. budget.add_num_batched_tokens(seq_group.request_id,
  496. num_running_tokens +
  497. negative_num_running_tokens)
  498. # OPTIMIZATION: Note that get_max_num_running_seqs is
  499. # expensive. For the default scheduling chase where
  500. # enable_chunking is False, num_seqs are updated before running
  501. # this method, so we don't have to update it again here.
  502. if enable_chunking:
  503. num_running_seqs = seq_group.get_max_num_running_seqs()
  504. budget.add_num_seqs(seq_group.request_id, num_running_seqs)
  505. if curr_loras is not None and seq_group.lora_int_id > 0:
  506. curr_loras.add(seq_group.lora_int_id)
  507. self._scheduler_running_outputs_cache.reset()
  508. self._scheduled_seq_group_cache.reset()
  509. return ret
  510. def _schedule_swapped(
  511. self,
  512. budget: SchedulingBudget,
  513. curr_loras: Optional[Set[int]],
  514. enable_chunking: bool = False,
  515. ) -> SchedulerSwappedInOutputs:
  516. """Schedule sequence groups that are swapped out.
  517. It schedules swapped requests as long as it fits `budget` and
  518. curr_loras <= max_lora from the scheduling config. The input arguments
  519. `budget` and `curr_loras` are updated based on scheduled seq_groups.
  520. Args:
  521. budget: The scheduling budget. The argument is in-place updated
  522. when any requests are swapped in.
  523. curr_loras: Currently batched lora request ids. The argument is
  524. in-place updated when any requests are swapped in.
  525. enable_chunking: If True, seq group can be chunked and only a
  526. chunked number of tokens are scheduled if
  527. `budget.num_batched_tokens` has not enough capacity to schedule
  528. all tokens.
  529. Returns:
  530. SchedulerSwappedInOutputs.
  531. """
  532. # Blocks that need to be swapped or copied before model execution.
  533. blocks_to_swap_in: List[Tuple[int, int]] = []
  534. blocks_to_copy: List[Tuple[int, int]] = []
  535. decode_seq_groups: List[ScheduledSequenceGroup] = []
  536. prefill_seq_groups: List[ScheduledSequenceGroup] = []
  537. infeasible_seq_groups: List[SequenceGroup] = []
  538. swapped_queue = self.swapped
  539. leftover_swapped: Deque[SequenceGroup] = deque()
  540. while swapped_queue:
  541. seq_group = swapped_queue[0]
  542. # If the sequence group cannot be swapped in, stop.
  543. is_prefill = seq_group.is_prefill()
  544. alloc_status = self.block_manager.can_swap_in(
  545. seq_group, self._get_num_lookahead_slots(is_prefill))
  546. if alloc_status == AllocStatus.LATER:
  547. break
  548. elif alloc_status == AllocStatus.NEVER:
  549. logger.warning(f"Failing the request {seq_group.request_id} "
  550. "because there's not enough kv cache blocks to "
  551. "run the entire sequence.")
  552. for seq in seq_group.get_seqs():
  553. seq.status = SequenceStatus.FINISHED_IGNORED
  554. infeasible_seq_groups.append(seq_group)
  555. swapped_queue.popleft()
  556. continue
  557. lora_int_id = 0
  558. if self.lora_enabled:
  559. lora_int_id = seq_group.lora_int_id
  560. assert curr_loras is not None
  561. assert self.lora_config is not None
  562. if (lora_int_id > 0 and (lora_int_id not in curr_loras)
  563. and len(curr_loras) >= self.lora_config.max_loras):
  564. # We don't have a space for another LoRA, so
  565. # we ignore this request for now.
  566. leftover_swapped.appendleft(seq_group)
  567. swapped_queue.popleft()
  568. continue
  569. # The total number of sequences in the RUNNING state should not
  570. # exceed the maximum number of sequences.
  571. num_new_seqs = seq_group.get_max_num_running_seqs()
  572. num_new_tokens, negative_num_new_tokens = self._get_num_new_tokens(
  573. seq_group, SequenceStatus.SWAPPED, enable_chunking, budget)
  574. if (num_new_tokens + negative_num_new_tokens == 0
  575. or not budget.can_schedule(num_new_tokens=num_new_tokens +
  576. negative_num_new_tokens,
  577. num_new_seqs=num_new_seqs)):
  578. break
  579. if lora_int_id > 0 and curr_loras is not None:
  580. curr_loras.add(lora_int_id)
  581. swapped_queue.popleft()
  582. self._swap_in(seq_group, blocks_to_swap_in)
  583. self._append_slots(seq_group, blocks_to_copy)
  584. is_prefill = seq_group.is_prefill()
  585. if is_prefill:
  586. prefill_seq_groups.append(
  587. ScheduledSequenceGroup(seq_group,
  588. token_chunk_size=num_new_tokens,
  589. negative_token_chunk_size=negative_num_new_tokens))
  590. else:
  591. decode_seq_groups.append(
  592. ScheduledSequenceGroup(seq_group, token_chunk_size=1,
  593. negative_token_chunk_size=1))
  594. budget.add_num_batched_tokens(seq_group.request_id,
  595. num_new_tokens +
  596. negative_num_new_tokens)
  597. budget.add_num_seqs(seq_group.request_id, num_new_seqs)
  598. swapped_queue.extendleft(leftover_swapped)
  599. return SchedulerSwappedInOutputs(
  600. decode_seq_groups=decode_seq_groups,
  601. prefill_seq_groups=prefill_seq_groups,
  602. blocks_to_swap_in=blocks_to_swap_in,
  603. blocks_to_copy=blocks_to_copy,
  604. num_lookahead_slots=self._get_num_lookahead_slots(
  605. is_prefill=False),
  606. infeasible_seq_groups=infeasible_seq_groups,
  607. )
  608. def _get_prompt_limit(self, seq_group: SequenceGroup) -> int:
  609. if self.scheduler_config.chunked_prefill_enabled:
  610. prompt_limit = self.scheduler_config.max_model_len
  611. else:
  612. prompt_limit = min(self.scheduler_config.max_model_len,
  613. self.scheduler_config.max_num_batched_tokens)
  614. # Model is fine tuned with long context. Return the fine tuned max_len.
  615. if (seq_group.lora_request
  616. and seq_group.lora_request.long_lora_max_len):
  617. assert prompt_limit <= seq_group.lora_request.long_lora_max_len
  618. return seq_group.lora_request.long_lora_max_len
  619. else:
  620. return prompt_limit
  621. def _schedule_prefills(
  622. self,
  623. budget: SchedulingBudget,
  624. curr_loras: Optional[Set[int]],
  625. enable_chunking: bool = False,
  626. ) -> SchedulerPrefillOutputs:
  627. """Schedule sequence groups that are in prefill stage.
  628. Note that the current scheduler treats PREEMPTED_FOR_RECOMPUTE
  629. as a new prefill (that starts from beginning -> most recently generated
  630. tokens).
  631. It schedules waiting requests as long as it fits `budget` and
  632. curr_loras <= max_lora from the scheduling config. The input arguments
  633. `budget` and `curr_loras` are updated based on scheduled seq_groups.
  634. Args:
  635. budget: The scheduling budget. The argument is in-place updated
  636. when any requests are scheduled.
  637. curr_loras: Currently batched lora request ids. The argument is
  638. in-place updated when any requests are scheduled.
  639. enable_chunking: If True, seq group can be chunked and only a
  640. chunked number of tokens are scheduled if
  641. `budget.num_batched_tokens` has not enough capacity to schedule
  642. all tokens.
  643. Returns:
  644. SchedulerPrefillOutputs.
  645. """
  646. ignored_seq_groups: List[SequenceGroup] = []
  647. seq_groups: List[SequenceGroup] = []
  648. waiting_queue = self.waiting
  649. leftover_waiting_sequences: Deque[SequenceGroup] = deque()
  650. while self._passed_delay(time.time()) and waiting_queue:
  651. seq_group = waiting_queue[0]
  652. waiting_seqs = seq_group.get_seqs(status=SequenceStatus.WAITING)
  653. assert len(waiting_seqs) == 1, (
  654. "Waiting sequence group should have only one prompt "
  655. "sequence.")
  656. num_new_tokens, negative_num_new_tokens = self._get_num_new_tokens(
  657. seq_group, SequenceStatus.WAITING, enable_chunking, budget)
  658. if not enable_chunking:
  659. num_prompt_tokens = waiting_seqs[0].get_len()
  660. assert num_new_tokens == num_prompt_tokens
  661. prompt_limit = self._get_prompt_limit(seq_group)
  662. if num_new_tokens + negative_num_new_tokens > prompt_limit:
  663. logger.warning(
  664. f"Input prompt ({num_new_tokens + negative_num_new_tokens} "
  665. "tokens) is too long"
  666. f" and exceeds limit of {prompt_limit}")
  667. for seq in waiting_seqs:
  668. seq.status = SequenceStatus.FINISHED_IGNORED
  669. ignored_seq_groups.append(seq_group)
  670. waiting_queue.popleft()
  671. continue
  672. # If the sequence group cannot be allocated, stop.
  673. can_allocate = self.block_manager.can_allocate(seq_group)
  674. if can_allocate == AllocStatus.LATER:
  675. break
  676. elif can_allocate == AllocStatus.NEVER:
  677. logger.warning(
  678. f"Input prompt ({num_new_tokens + negative_num_new_tokens} "
  679. "tokens) is too long"
  680. " and exceeds the capacity of block_manager")
  681. for seq in waiting_seqs:
  682. seq.status = SequenceStatus.FINISHED_IGNORED
  683. ignored_seq_groups.append(seq_group)
  684. waiting_queue.popleft()
  685. continue
  686. lora_int_id = 0
  687. if self.lora_enabled:
  688. lora_int_id = seq_group.lora_int_id
  689. assert curr_loras is not None
  690. assert self.lora_config is not None
  691. if (self.lora_enabled and lora_int_id > 0
  692. and lora_int_id not in curr_loras
  693. and len(curr_loras) >= self.lora_config.max_loras):
  694. # We don't have a space for another LoRA, so
  695. # we ignore this request for now.
  696. leftover_waiting_sequences.appendleft(seq_group)
  697. waiting_queue.popleft()
  698. continue
  699. num_new_seqs = seq_group.get_max_num_running_seqs()
  700. if (num_new_tokens + negative_num_new_tokens == 0
  701. or not budget.can_schedule(num_new_tokens=num_new_tokens +
  702. negative_num_new_tokens,
  703. num_new_seqs=num_new_seqs)):
  704. break
  705. # Can schedule this request.
  706. if curr_loras is not None and lora_int_id > 0:
  707. curr_loras.add(lora_int_id)
  708. waiting_queue.popleft()
  709. self._allocate_and_set_running(seq_group)
  710. seq_group.init_multi_step(
  711. num_scheduler_steps=self._get_num_lookahead_slots(
  712. is_prefill=True) + 1)
  713. seq_groups.append(
  714. ScheduledSequenceGroup(seq_group=seq_group,
  715. token_chunk_size=num_new_tokens,
  716. negative_token_chunk_size=negative_num_new_tokens))
  717. budget.add_num_batched_tokens(seq_group.request_id,
  718. num_new_tokens +
  719. negative_num_new_tokens)
  720. budget.add_num_seqs(seq_group.request_id, num_new_seqs)
  721. # Queue requests that couldn't be scheduled.
  722. waiting_queue.extendleft(leftover_waiting_sequences)
  723. if len(seq_groups) > 0:
  724. self.prev_prompt = True
  725. return SchedulerPrefillOutputs(
  726. seq_groups=seq_groups,
  727. ignored_seq_groups=ignored_seq_groups,
  728. num_lookahead_slots=self._get_num_lookahead_slots(is_prefill=True))
  729. def _schedule_default(self) -> SchedulerOutputs:
  730. """Schedule queued requests.
  731. The current policy is designed to optimize the throughput. First,
  732. it batches as many prefill requests as possible. And it schedules
  733. decodes. If there's a pressure on GPU memory, decode requests can
  734. be swapped or preempted.
  735. """
  736. # Include running requests to the budget.
  737. budget = SchedulingBudget(
  738. token_budget=self.scheduler_config.max_num_batched_tokens,
  739. max_num_seqs=self.scheduler_config.max_num_seqs,
  740. )
  741. # Make sure we include num running seqs before scheduling prefill,
  742. # so that we don't schedule beyond max_num_seqs for prefill.
  743. for seq_group in self.running:
  744. budget.add_num_seqs(seq_group.request_id,
  745. seq_group.get_max_num_running_seqs())
  746. curr_loras = set(
  747. seq_group.lora_int_id for seq_group in self.running
  748. if seq_group.lora_int_id > 0) if self.lora_enabled else None
  749. prefills = SchedulerPrefillOutputs.create_empty()
  750. running_scheduled = SchedulerRunningOutputs.create_empty()
  751. swapped_in = SchedulerSwappedInOutputs.create_empty()
  752. # If any requests are swapped, prioritized swapped requests.
  753. if not self.swapped:
  754. prefills = self._schedule_prefills(budget,
  755. curr_loras,
  756. enable_chunking=False)
  757. # Don't schedule decodes if prefills are scheduled.
  758. # NOTE: If `_schedule_prefills` doesn't enable chunking, self.running
  759. # only contains decode requests, not chunked prefills.
  760. if len(prefills.seq_groups) == 0:
  761. running_scheduled = self._schedule_running(budget,
  762. curr_loras,
  763. enable_chunking=False)
  764. # If any sequence group is preempted, do not swap in any sequence
  765. # group. because it means there's no slot for new running requests.
  766. if len(running_scheduled.preempted) + len(
  767. running_scheduled.swapped_out) == 0:
  768. swapped_in = self._schedule_swapped(budget, curr_loras)
  769. assert (budget.num_batched_tokens <=
  770. self.scheduler_config.max_num_batched_tokens)
  771. assert budget.num_curr_seqs <= self.scheduler_config.max_num_seqs
  772. # Update waiting requests.
  773. self.waiting.extendleft(running_scheduled.preempted)
  774. # Update new running requests.
  775. if len(prefills.seq_groups) > 0:
  776. self.running.extend([s.seq_group for s in prefills.seq_groups])
  777. self.running.extend(running_scheduled.decode_seq_groups_list)
  778. if len(swapped_in.decode_seq_groups) > 0:
  779. self.running.extend(
  780. [s.seq_group for s in swapped_in.decode_seq_groups])
  781. # Update swapped requests.
  782. self.swapped.extend(running_scheduled.swapped_out)
  783. preempted = (len(running_scheduled.preempted) +
  784. len(running_scheduled.swapped_out))
  785. # There should be no prefill from running queue because this policy
  786. # doesn't allow chunked prefills.
  787. assert len(running_scheduled.prefill_seq_groups) == 0
  788. assert len(swapped_in.prefill_seq_groups) == 0
  789. # Merge lists
  790. num_prefill_groups = len(prefills.seq_groups)
  791. if num_prefill_groups > 0:
  792. scheduled_seq_groups = prefills.seq_groups
  793. scheduled_seq_groups.extend(running_scheduled.decode_seq_groups)
  794. else:
  795. scheduled_seq_groups = running_scheduled.decode_seq_groups
  796. scheduled_seq_groups.extend(swapped_in.decode_seq_groups)
  797. blocks_to_copy = running_scheduled.blocks_to_copy
  798. blocks_to_copy.extend(swapped_in.blocks_to_copy)
  799. ignored_seq_groups = prefills.ignored_seq_groups
  800. ignored_seq_groups.extend(swapped_in.infeasible_seq_groups)
  801. return SchedulerOutputs(
  802. scheduled_seq_groups=scheduled_seq_groups,
  803. num_prefill_groups=num_prefill_groups,
  804. num_batched_tokens=budget.num_batched_tokens,
  805. blocks_to_swap_in=swapped_in.blocks_to_swap_in,
  806. blocks_to_swap_out=running_scheduled.blocks_to_swap_out,
  807. blocks_to_copy=blocks_to_copy,
  808. ignored_seq_groups=ignored_seq_groups,
  809. num_lookahead_slots=running_scheduled.num_lookahead_slots,
  810. running_queue_size=len(self.running),
  811. preempted=preempted,
  812. )
  813. def _schedule_chunked_prefill(self) -> SchedulerOutputs:
  814. """Schedule queued requests.
  815. Chunked prefill allows to chunk prefill requests, batch them together
  816. with decode requests. This policy 1. schedule as many decoding requests
  817. as possible. 2. schedule chunked prefill requests that are not
  818. finished. 3. schedule swapped request. 4. schedule new prefill
  819. requests.
  820. The policy can sustain the high GPU utilization because it can put
  821. prefill and decodes requests to the same batch, while it improves
  822. inter token latency because decodes requests don't need to be blocked
  823. by prefill requests.
  824. """
  825. budget = SchedulingBudget(
  826. token_budget=self.scheduler_config.max_num_batched_tokens,
  827. max_num_seqs=self.scheduler_config.max_num_seqs,
  828. )
  829. curr_loras: Set[int] = set()
  830. prefills = SchedulerPrefillOutputs.create_empty()
  831. swapped_in = SchedulerSwappedInOutputs.create_empty()
  832. # Decoding should be always scheduled first by fcfs.
  833. running_scheduled = self._schedule_running(budget,
  834. curr_loras,
  835. enable_chunking=True)
  836. # Schedule swapped out requests.
  837. # If preemption happens, it means we don't have space for swap-in.
  838. if len(running_scheduled.preempted) + len(
  839. running_scheduled.swapped_out) == 0:
  840. swapped_in = self._schedule_swapped(budget, curr_loras)
  841. # Schedule new prefills.
  842. prefills = self._schedule_prefills(budget,
  843. curr_loras,
  844. enable_chunking=True)
  845. assert (budget.num_batched_tokens <=
  846. self.scheduler_config.max_num_batched_tokens)
  847. assert budget.num_curr_seqs <= self.scheduler_config.max_num_seqs
  848. # Update waiting requests.
  849. self.waiting.extendleft(running_scheduled.preempted)
  850. # Update new running requests.
  851. self.running.extend([s.seq_group for s in prefills.seq_groups])
  852. self.running.extend(
  853. [s.seq_group for s in running_scheduled.decode_seq_groups])
  854. self.running.extend(
  855. [s.seq_group for s in running_scheduled.prefill_seq_groups])
  856. self.running.extend(
  857. [s.seq_group for s in swapped_in.decode_seq_groups])
  858. self.running.extend(
  859. [s.seq_group for s in swapped_in.prefill_seq_groups])
  860. # Update swapped requests.
  861. self.swapped.extend(running_scheduled.swapped_out)
  862. return SchedulerOutputs(
  863. scheduled_seq_groups=(prefills.seq_groups +
  864. running_scheduled.prefill_seq_groups +
  865. swapped_in.prefill_seq_groups +
  866. running_scheduled.decode_seq_groups +
  867. swapped_in.decode_seq_groups),
  868. num_prefill_groups=(len(prefills.seq_groups) +
  869. len(swapped_in.prefill_seq_groups) +
  870. len(running_scheduled.prefill_seq_groups)),
  871. num_batched_tokens=budget.num_batched_tokens,
  872. blocks_to_swap_in=swapped_in.blocks_to_swap_in,
  873. blocks_to_swap_out=running_scheduled.blocks_to_swap_out,
  874. blocks_to_copy=running_scheduled.blocks_to_copy +
  875. swapped_in.blocks_to_copy,
  876. ignored_seq_groups=prefills.ignored_seq_groups +
  877. swapped_in.infeasible_seq_groups,
  878. num_lookahead_slots=running_scheduled.num_lookahead_slots,
  879. running_queue_size=len(self.running),
  880. preempted=(len(running_scheduled.preempted) +
  881. len(running_scheduled.swapped_out)),
  882. )
  883. def _schedule(self) -> SchedulerOutputs:
  884. """Schedule queued requests."""
  885. if self.scheduler_config.chunked_prefill_enabled:
  886. return self._schedule_chunked_prefill()
  887. else:
  888. return self._schedule_default()
  889. def _can_append_slots(self, seq_group: SequenceGroup) -> bool:
  890. """Determine whether or not we have enough space in the KV cache to
  891. continue generation of the sequence group.
  892. """
  893. # It is True only for testing case to trigger artificial preemption.
  894. if (self.enable_artificial_preemption
  895. and random.uniform(0, 1) < ARTIFICIAL_PREEMPTION_PROB
  896. and self.artificial_preempt_cnt > 0):
  897. self.artificial_preempt_cnt -= 1
  898. return False
  899. # Appending slots only occurs in decoding.
  900. is_prefill = False
  901. return self.block_manager.can_append_slots(
  902. seq_group=seq_group,
  903. num_lookahead_slots=self._get_num_lookahead_slots(is_prefill),
  904. )
  905. def schedule(self) -> Tuple[List[SequenceGroupMetadata], SchedulerOutputs]:
  906. # Schedule sequence groups.
  907. # This function call changes the internal states of the scheduler
  908. # such as self.running, self.swapped, and self.waiting.
  909. scheduler_outputs = self._schedule()
  910. now = time.time()
  911. if not self.cache_config.enable_prefix_caching:
  912. common_computed_block_nums = []
  913. # Create input data structures.
  914. seq_group_metadata_list: List[SequenceGroupMetadata] = []
  915. for i, scheduled_seq_group in enumerate(
  916. scheduler_outputs.scheduled_seq_groups):
  917. seq_group = scheduled_seq_group.seq_group
  918. token_chunk_size = scheduled_seq_group.token_chunk_size
  919. negative_token_chunk_size = \
  920. scheduled_seq_group.negative_token_chunk_size
  921. seq_group.maybe_set_first_scheduled_time(now)
  922. # seq_id -> SequenceData
  923. seq_data: Dict[int, SequenceData] = {}
  924. # seq_id -> physical block numbers
  925. block_tables: Dict[int, List[int]] = {}
  926. if seq_group.is_encoder_decoder():
  927. # Encoder associated with SequenceGroup
  928. encoder_seq_data = seq_group.get_encoder_seq().data
  929. # Block table for cross-attention
  930. # Also managed at SequenceGroup level
  931. cross_block_table = self.block_manager.get_cross_block_table(
  932. seq_group)
  933. else:
  934. encoder_seq_data = None
  935. cross_block_table = None
  936. if seq_group.has_negative_seqs():
  937. # seq_id -> SequenceData
  938. negative_seq_data: Dict[int, SequenceData] = {}
  939. # seq_id -> physical block numbers
  940. negative_block_tables: Dict[int, List[int]] = {}
  941. else:
  942. negative_seq_data = None
  943. negative_block_tables = None
  944. for seq in seq_group.get_seqs(status=SequenceStatus.RUNNING):
  945. seq_id = seq.seq_id
  946. seq_data[seq_id] = seq.data
  947. block_tables[seq_id] = self.block_manager.get_block_table(seq)
  948. self.block_manager.access_all_blocks_in_seq(seq, now)
  949. if seq_group.has_negative_seqs():
  950. negative_seq = seq_group.negative_seqs_dict[seq_id]
  951. negative_seq_data[seq_id] = negative_seq.data
  952. negative_block_tables[seq_id] = \
  953. self.block_manager.get_negative_block_table(negative_seq)
  954. if self.cache_config.enable_prefix_caching:
  955. common_computed_block_nums = (
  956. self.block_manager.get_common_computed_block_ids(
  957. seq_group.get_seqs(status=SequenceStatus.RUNNING)))
  958. do_sample = True
  959. is_prompt = seq_group.is_prefill()
  960. # We should send the metadata to workers when the first prefill
  961. # is sent. Subsequent requests could be chunked prefill or decode.
  962. is_first_prefill = False
  963. if is_prompt:
  964. seqs = seq_group.get_seqs()
  965. # Prefill has only 1 sequence.
  966. assert len(seqs) == 1
  967. num_computed_tokens = seqs[0].data.get_num_computed_tokens()
  968. is_first_prefill = num_computed_tokens == 0
  969. # In the next iteration, all prompt tokens are not computed.
  970. # It means the prefill is chunked, and we don't need sampling.
  971. # NOTE: We use get_len instead of get_prompt_len because when
  972. # a sequence is preempted, prefill includes previous generated
  973. # output tokens.
  974. if (token_chunk_size + num_computed_tokens <
  975. seqs[0].data.get_len()):
  976. do_sample = False
  977. negative_seqs = seq_group.get_negative_seqs()
  978. assert len(negative_seqs) == 1
  979. negative_num_computed_tokens = \
  980. negative_seqs[0].data.get_num_computed_tokens()
  981. if (negative_token_chunk_size + negative_num_computed_tokens <
  982. negative_seqs[0].data.get_len()):
  983. do_sample = False
  984. # It assumes the scheduled_seq_groups is ordered by
  985. # prefill < decoding.
  986. if is_first_prefill or not self.scheduler_config.send_delta_data:
  987. seq_group_metadata = SequenceGroupMetadata(
  988. request_id=seq_group.request_id,
  989. is_prompt=is_prompt,
  990. seq_data=seq_data,
  991. sampling_params=seq_group.sampling_params,
  992. block_tables=block_tables,
  993. do_sample=do_sample,
  994. pooling_params=seq_group.pooling_params,
  995. token_chunk_size=token_chunk_size,
  996. negative_token_chunk_size=negative_token_chunk_size,
  997. lora_request=seq_group.lora_request,
  998. computed_block_nums=common_computed_block_nums,
  999. encoder_seq_data=encoder_seq_data,
  1000. cross_block_table=cross_block_table,
  1001. negative_seq_data=negative_seq_data,
  1002. negative_block_tables=negative_block_tables,
  1003. state=seq_group.state,
  1004. # `multi_modal_data` will only be present for the 1st comm
  1005. # between engine and worker.
  1006. # the subsequent comms can still use delta, but
  1007. # `multi_modal_data` will be None.
  1008. multi_modal_data=seq_group.multi_modal_data
  1009. if scheduler_outputs.num_prefill_groups > 0 else None,
  1010. prompt_adapter_request=seq_group.prompt_adapter_request,
  1011. )
  1012. else:
  1013. # When SPMD mode is enabled, we only send delta data except for
  1014. # the first request to reduce serialization cost.
  1015. seq_data_delta = {}
  1016. for id, data in seq_data.items():
  1017. seq_data_delta[id] = data.get_delta_and_reset()
  1018. seq_group_metadata = SequenceGroupMetadataDelta(
  1019. seq_data_delta,
  1020. seq_group.request_id,
  1021. block_tables,
  1022. is_prompt,
  1023. do_sample=do_sample,
  1024. token_chunk_size=token_chunk_size,
  1025. computed_block_nums=common_computed_block_nums,
  1026. )
  1027. seq_group_metadata_list.append(seq_group_metadata)
  1028. # Now that the batch has been created, we can assume all blocks in the
  1029. # batch will have been computed before the next scheduling invocation.
  1030. # This is because the engine assumes that a failure in model execution
  1031. # will crash the Aphrodite instance / will not retry.
  1032. for scheduled_seq_group in scheduler_outputs.scheduled_seq_groups:
  1033. self.block_manager.mark_blocks_as_computed(
  1034. scheduled_seq_group.seq_group)
  1035. return seq_group_metadata_list, scheduler_outputs
  1036. def fork_seq(self, parent_seq: Sequence, child_seq: Sequence) -> None:
  1037. self.block_manager.fork(parent_seq, child_seq)
  1038. def free_seq(self, seq: Sequence) -> None:
  1039. """Free a sequence from a block table."""
  1040. self.block_manager.free(seq)
  1041. def free_finished_seq_groups(self) -> None:
  1042. remaining: Deque[SequenceGroup] = deque()
  1043. for seq_group in self.running:
  1044. if seq_group.is_finished():
  1045. # Free cross-attention block table, if it exists
  1046. self._free_seq_group_cross_attn_blocks(seq_group)
  1047. # Add the finished requests to the finished requests list.
  1048. # This list will be used to update the Mamba cache in the
  1049. # next step.
  1050. self._finished_requests_ids.append(seq_group.request_id)
  1051. else:
  1052. remaining.append(seq_group)
  1053. self.running = remaining
  1054. def _allocate_and_set_running(self, seq_group: SequenceGroup) -> None:
  1055. self.block_manager.allocate(seq_group)
  1056. for seq in seq_group.get_seqs(status=SequenceStatus.WAITING):
  1057. seq.status = SequenceStatus.RUNNING
  1058. def _append_slots(
  1059. self,
  1060. seq_group: SequenceGroup,
  1061. blocks_to_copy: List[Tuple[int, int]],
  1062. ) -> None:
  1063. """Appends new slots to the sequences in the given sequence group.
  1064. Args:
  1065. seq_group (SequenceGroup): The sequence group containing the
  1066. sequences to append slots to.
  1067. blocks_to_copy (List[Tuple[int, int]]): A list of tuple of two
  1068. ints, the first int is the source block index, and the second
  1069. int is the destination block index. This list is updated with
  1070. the new source and destination block indices for the appended
  1071. slots.
  1072. """
  1073. num_lookahead_slots = self._get_num_lookahead_slots(is_prefill=False)
  1074. seq_group.init_multi_step(num_scheduler_steps=num_lookahead_slots + 1)
  1075. for seq in seq_group.get_seqs(status=SequenceStatus.RUNNING):
  1076. cows = self.block_manager.append_slots(seq, num_lookahead_slots,
  1077. seq_group)
  1078. assert len(cows) == 0
  1079. if len(cows) > 0:
  1080. blocks_to_copy.extend(cows)
  1081. def _preempt(
  1082. self,
  1083. seq_group: SequenceGroup,
  1084. blocks_to_swap_out: List[Tuple[int, int]],
  1085. preemption_mode: Optional[PreemptionMode] = None,
  1086. ) -> PreemptionMode:
  1087. # If preemption mode is not specified, we determine the mode as follows:
  1088. # We use recomputation by default since it incurs lower overhead than
  1089. # swapping. However, when the sequence group has multiple sequences
  1090. # (e.g., beam search), recomputation is not currently supported. In
  1091. # such a case, we use swapping instead.
  1092. # FIXME: This makes our scheduling policy a bit bizarre.
  1093. # As swapped sequences are prioritized over waiting sequences,
  1094. # sequence groups with multiple sequences are implicitly prioritized
  1095. # over sequence groups with a single sequence.
  1096. # TODO: Support recomputation for sequence groups with multiple
  1097. # sequences. This may require a more sophisticated CUDA kernel.
  1098. if self.user_specified_preemption_mode is None:
  1099. if seq_group.get_max_num_running_seqs() == 1:
  1100. preemption_mode = PreemptionMode.RECOMPUTE
  1101. else:
  1102. preemption_mode = PreemptionMode.SWAP
  1103. elif self.user_specified_preemption_mode == "swap":
  1104. preemption_mode = PreemptionMode.SWAP
  1105. else:
  1106. preemption_mode = PreemptionMode.RECOMPUTE
  1107. if self.num_cumulative_preemption % 50 == 0:
  1108. logger.warning(
  1109. f"Sequence group {seq_group.request_id} is preempted by "
  1110. f"{preemption_mode} mode because there is "
  1111. "not enough KV cache space. This can affect the end-to-end "
  1112. "performance. Increase gpu_memory_utilization or "
  1113. "tensor_parallel_size to provide more KV cache memory. "
  1114. "total_num_cumulative_preemption="
  1115. f"{self.num_cumulative_preemption + 1}")
  1116. self.num_cumulative_preemption += 1
  1117. if preemption_mode == PreemptionMode.RECOMPUTE:
  1118. self._preempt_by_recompute(seq_group)
  1119. elif preemption_mode == PreemptionMode.SWAP:
  1120. self._preempt_by_swap(seq_group, blocks_to_swap_out)
  1121. else:
  1122. raise AssertionError("Invalid preemption mode.")
  1123. return preemption_mode
  1124. def _preempt_by_recompute(
  1125. self,
  1126. seq_group: SequenceGroup,
  1127. ) -> None:
  1128. seqs = seq_group.get_seqs(status=SequenceStatus.RUNNING)
  1129. assert len(seqs) == 1
  1130. for seq in seqs:
  1131. seq.status = SequenceStatus.WAITING
  1132. self.free_seq(seq)
  1133. seq.reset_state_for_recompute()
  1134. negative_seq = seq_group.negative_seqs_dict[seq.seq_id]
  1135. negative_seq.reset_state_for_recompute()
  1136. def _preempt_by_swap(
  1137. self,
  1138. seq_group: SequenceGroup,
  1139. blocks_to_swap_out: List[Tuple[int, int]],
  1140. ) -> None:
  1141. self._swap_out(seq_group, blocks_to_swap_out)
  1142. def _swap_in(
  1143. self,
  1144. seq_group: SequenceGroup,
  1145. blocks_to_swap_in: List[Tuple[int, int]],
  1146. ) -> None:
  1147. mapping = self.block_manager.swap_in(seq_group)
  1148. blocks_to_swap_in.extend(mapping)
  1149. for seq in seq_group.get_seqs(status=SequenceStatus.SWAPPED):
  1150. seq.status = SequenceStatus.RUNNING
  1151. def _swap_out(
  1152. self,
  1153. seq_group: SequenceGroup,
  1154. blocks_to_swap_out: List[Tuple[int, int]],
  1155. ) -> None:
  1156. if not self.block_manager.can_swap_out(seq_group):
  1157. # FIXME: Abort the sequence group instead of aborting the
  1158. # entire engine.
  1159. raise RuntimeError(
  1160. "Aborted due to the lack of CPU swap space. Please increase "
  1161. "the swap space to avoid this error.")
  1162. mapping = self.block_manager.swap_out(seq_group)
  1163. blocks_to_swap_out.extend(mapping)
  1164. for seq in seq_group.get_seqs(status=SequenceStatus.RUNNING):
  1165. seq.status = SequenceStatus.SWAPPED
  1166. def _passed_delay(self, now: float) -> bool:
  1167. if self.prev_prompt:
  1168. self.last_prompt_latency = now - self.prev_time
  1169. self.prev_time, self.prev_prompt = now, False
  1170. # Delay scheduling prompts to let waiting queue fill up
  1171. if self.scheduler_config.delay_factor > 0 and self.waiting:
  1172. earliest_arrival_time = min(
  1173. [e.metrics.arrival_time for e in self.waiting])
  1174. passed_delay = (
  1175. (now - earliest_arrival_time) >
  1176. (self.scheduler_config.delay_factor * self.last_prompt_latency)
  1177. or not self.running)
  1178. else:
  1179. passed_delay = True
  1180. return passed_delay
  1181. def _get_num_lookahead_slots(self, is_prefill: bool) -> int:
  1182. """The number of slots to allocate per sequence per step, beyond known
  1183. token ids. Speculative decoding uses these slots to store KV activations
  1184. of tokens which may or may not be accepted.
  1185. Speculative decoding does not yet support prefill, so we do not perform
  1186. lookahead allocation for prefill.
  1187. """
  1188. if is_prefill:
  1189. return 0
  1190. return self.scheduler_config.num_lookahead_slots
  1191. def _get_num_new_tokens(self, seq_group: SequenceGroup,
  1192. status: SequenceStatus, enable_chunking: bool,
  1193. budget: SchedulingBudget) -> int:
  1194. """Get the next new tokens to compute for a given sequence group
  1195. that's in a given `status`.
  1196. The API could chunk the number of tokens to compute based on `budget`
  1197. if `enable_chunking` is True. If a sequence group has multiple
  1198. sequences (e.g., running beam search), it means it is in decoding
  1199. phase, so chunking doesn't happen.
  1200. Returns 0 if the new token cannot be computed due to token budget.
  1201. """
  1202. num_new_tokens = 0
  1203. negative_num_new_tokens = 0
  1204. seqs = seq_group.get_seqs(status=status)
  1205. for seq in seqs:
  1206. num_new_tokens += seq.get_num_new_tokens()
  1207. if seq_group.has_negative_seqs():
  1208. negative_seq = seq_group.negative_seqs_dict[seq.seq_id]
  1209. negative_num_new_tokens += negative_seq.get_num_new_tokens()
  1210. assert num_new_tokens + negative_num_new_tokens > 0
  1211. # Chunk if a running request cannot fit in.
  1212. # If number of seq > 1, it means it is doing beam search in a
  1213. # decode phase. Do not chunk in that case.
  1214. if seq_group.has_negative_seqs():
  1215. if enable_chunking and len(seqs) == 1:
  1216. if num_new_tokens + negative_num_new_tokens < \
  1217. budget.remaining_token_budget():
  1218. return num_new_tokens, negative_num_new_tokens
  1219. else:
  1220. half_remaining_budget = budget.remaining_token_budget() // 2
  1221. rt_num_new_tokens = min(num_new_tokens,
  1222. half_remaining_budget)
  1223. rt_negative_num_new_tokens = min(negative_num_new_tokens,
  1224. half_remaining_budget)
  1225. # ruff: noqa
  1226. if rt_num_new_tokens == num_new_tokens and \
  1227. rt_negative_num_new_tokens != negative_num_new_tokens:
  1228. return 0, 0
  1229. elif rt_num_new_tokens != num_new_tokens and \
  1230. rt_negative_num_new_tokens == negative_num_new_tokens:
  1231. return 0, 0
  1232. else:
  1233. return rt_num_new_tokens, rt_negative_num_new_tokens
  1234. return num_new_tokens, negative_num_new_tokens
  1235. else:
  1236. if enable_chunking and len(seqs) == 1:
  1237. num_new_tokens = min(num_new_tokens,
  1238. budget.remaining_token_budget())
  1239. return num_new_tokens, 0