@@ -141,9 +141,10 @@ template <ggml_type type, int ncols_dst>
141
141
__launch_bounds__ (calc_nwarps(ncols_dst, get_device_table_id())*ggml_cuda_get_physical_warp_size(), 1)
142
142
static __global__ void mul_mat_vec_q(
143
143
const void * __restrict__ vx, const void * __restrict__ vy, const int32_t * __restrict__ ids, float * __restrict__ dst,
144
- const int ncols_x, const int nchannels_y, const int stride_row_x, const int stride_col_y, const int stride_col_dst,
145
- const int channel_ratio, const int stride_channel_x, const int stride_channel_y, const int stride_channel_dst,
146
- const int sample_ratio, const int stride_sample_x, const int stride_sample_y, const int stride_sample_dst) {
144
+ const uint32_t ncols_x, const uint32_t nchannels_y, const uint32_t stride_row_x, const uint32_t stride_col_y,
145
+ const uint32_t stride_col_dst, const uint3 channel_ratio, const uint32_t stride_channel_x,
146
+ const uint32_t stride_channel_y, const uint32_t stride_channel_dst, const uint3 sample_ratio,
147
+ const uint32_t stride_sample_x, const uint32_t stride_sample_y, const uint32_t stride_sample_dst) {
147
148
148
149
constexpr int qk = ggml_cuda_type_traits<type>::qk;
149
150
constexpr int qi = ggml_cuda_type_traits<type>::qi;
@@ -161,12 +162,12 @@ static __global__ void mul_mat_vec_q(
161
162
constexpr int blocks_per_iter = vdr * nwarps*warp_size / qi;
162
163
163
164
// The MUL_MAT_ID code path with ids != nullptr is only implemented for ncols_dst == 1.
164
- const int channel_dst = blockIdx .y ;
165
- const int channel_x = ncols_dst == 1 && ids ? ids[channel_dst] : channel_dst / channel_ratio;
166
- const int channel_y = ncols_dst == 1 && ids ? channel_dst % nchannels_y : channel_dst;
167
- const int sample_dst = blockIdx .z ;
168
- const int sample_x = sample_dst / sample_ratio;
169
- const int sample_y = sample_dst;
165
+ const uint32_t channel_dst = blockIdx .y ;
166
+ const uint32_t channel_x = ncols_dst == 1 && ids ? ids[channel_dst] : fastdiv ( channel_dst, channel_ratio) ;
167
+ const uint32_t channel_y = ncols_dst == 1 && ids ? channel_dst % nchannels_y : channel_dst;
168
+ const uint32_t sample_dst = blockIdx .z ;
169
+ const uint32_t sample_x = fastdiv ( sample_dst, sample_ratio) ;
170
+ const uint32_t sample_y = sample_dst;
170
171
171
172
// partial sum for each thread
172
173
float tmp[ncols_dst][rows_per_cuda_block] = {{0 .0f }};
@@ -247,95 +248,79 @@ static void mul_mat_vec_q_switch_ncols_dst(
247
248
GGML_ASSERT (ncols_x % ggml_blck_size (type) == 0 );
248
249
GGML_ASSERT (ncols_dst <= MMVQ_MAX_BATCH_SIZE);
249
250
250
- const int channel_ratio = nchannels_dst / nchannels_x;
251
- const int sample_ratio = nsamples_dst / nsamples_x;
251
+ const uint3 channel_ratio = init_fastdiv_values ( nchannels_dst / nchannels_x) ;
252
+ const uint3 sample_ratio = init_fastdiv_values ( nsamples_dst / nsamples_x) ;
252
253
253
254
const int device = ggml_cuda_get_device ();
254
255
const int warp_size = ggml_cuda_info ().devices [device].warp_size ;
255
256
const mmvq_parameter_table_id table_id = get_device_table_id (ggml_cuda_info ().devices [device].cc );
256
257
257
258
GGML_ASSERT (!ids || ncols_dst == 1 );
258
259
switch (ncols_dst) {
259
- case 1 :
260
- {
260
+ case 1 : {
261
261
constexpr int c_ncols_dst = 1 ;
262
262
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
263
263
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
264
264
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
265
265
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
266
266
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
267
- break ;
268
- }
269
- case 2 :
270
- {
267
+ } break ;
268
+ case 2 : {
271
269
constexpr int c_ncols_dst = 2 ;
272
270
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
273
271
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
274
272
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
275
273
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
276
274
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
277
- break ;
278
- }
279
- case 3 :
280
- {
275
+ } break ;
276
+ case 3 : {
281
277
constexpr int c_ncols_dst = 3 ;
282
278
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
283
279
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
284
280
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
285
281
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
286
282
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
287
- break ;
288
- }
289
- case 4 :
290
- {
283
+ } break ;
284
+ case 4 : {
291
285
constexpr int c_ncols_dst = 4 ;
292
286
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
293
287
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
294
288
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
295
289
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
296
290
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
297
- break ;
298
- }
299
- case 5 :
300
- {
291
+ } break ;
292
+ case 5 : {
301
293
constexpr int c_ncols_dst = 5 ;
302
294
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
303
295
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
304
296
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
305
297
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
306
298
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
307
- break ;
308
- }
309
- case 6 :
310
- {
299
+ } break ;
300
+ case 6 : {
311
301
constexpr int c_ncols_dst = 6 ;
312
302
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
313
303
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
314
304
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
315
305
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
316
306
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
317
- break ;
318
- }
319
- case 7 :
320
- {
307
+ } break ;
308
+ case 7 : {
321
309
constexpr int c_ncols_dst = 7 ;
322
310
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
323
311
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
324
312
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
325
313
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
326
314
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
327
- break ;
328
- }
329
- case 8 :
330
- {
315
+ } break ;
316
+ case 8 : {
331
317
constexpr int c_ncols_dst = 8 ;
332
318
std::pair<dim3 , dim3 > dims = calc_launch_params (c_ncols_dst, nrows_x, nchannels_dst, nsamples_dst, warp_size, table_id);
333
319
mul_mat_vec_q<type, c_ncols_dst><<<dims.first, dims.second, 0 , stream>>>
334
320
(vx, vy, ids, dst, ncols_x, nchannels_y, stride_row_x, stride_col_y, stride_col_dst,
335
321
channel_ratio, stride_channel_x, stride_channel_y, stride_channel_dst,
336
322
sample_ratio, stride_sample_x, stride_sample_y, stride_sample_dst);
337
- break ;
338
- }
323
+ } break ;
339
324
default :
340
325
GGML_ABORT (" fatal error" );
341
326
break ;
0 commit comments