mirror of
				https://github.com/ggml-org/llama.cpp.git
				synced 2025-10-30 08:42:00 +00:00 
			
		
		
		
	cuda : fix dmmv cols requirement to 2*GGML_CUDA_DMMV_X (#8800)
* cuda : fix dmmv cols requirement to 2*GGML_CUDA_DMMV_X * update asserts * only use dmmv for supported types * add test
This commit is contained in:
		| @@ -1885,10 +1885,9 @@ static void ggml_cuda_mul_mat_batched_cublas(ggml_backend_cuda_context & ctx, co | ||||
| static void ggml_cuda_mul_mat(ggml_backend_cuda_context & ctx, const ggml_tensor * src0, const ggml_tensor * src1, ggml_tensor * dst) { | ||||
|     const bool split = ggml_backend_buffer_is_cuda_split(src0->buffer); | ||||
|  | ||||
|     bool use_dequantize_mul_mat_vec = (ggml_is_quantized(src0->type) || src0->type == GGML_TYPE_F16) | ||||
|     bool use_dequantize_mul_mat_vec = ggml_cuda_dmmv_type_supported(src0->type) | ||||
|         && src1->type == GGML_TYPE_F32 && dst->type == GGML_TYPE_F32 | ||||
|         && src0->ne[0] % GGML_CUDA_DMMV_X == 0 && src0->ne[0] >= GGML_CUDA_DMMV_X*2 | ||||
|         && src1->ne[1] == 1; | ||||
|         && src0->ne[0] % (GGML_CUDA_DMMV_X*2) == 0 && src1->ne[1] == 1; | ||||
|     bool          use_mul_mat_vec_q =  ggml_is_quantized(src0->type) | ||||
|         && src1->type == GGML_TYPE_F32 && dst->type == GGML_TYPE_F32 | ||||
|         && src1->ne[1] <= MMVQ_MAX_BATCH_SIZE; | ||||
|   | ||||
| @@ -500,7 +500,7 @@ static __global__ void dequantize_mul_mat_vec(const void * __restrict__ vx, cons | ||||
| } | ||||
|  | ||||
| static void dequantize_mul_mat_vec_q4_0_cuda(const void * vx, const dfloat * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) { | ||||
|     GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0); | ||||
|     GGML_ASSERT(ncols % (GGML_CUDA_DMMV_X*2) == 0); | ||||
|     const int block_num_y = (nrows + GGML_CUDA_MMV_Y - 1) / GGML_CUDA_MMV_Y; | ||||
|     // the number of rows may exceed maximum grid size in the y or z dimensions, use the x dimension instead | ||||
|     const dim3 block_nums(block_num_y, 1, 1); | ||||
| @@ -510,7 +510,7 @@ static void dequantize_mul_mat_vec_q4_0_cuda(const void * vx, const dfloat * y, | ||||
| } | ||||
|  | ||||
| static void dequantize_mul_mat_vec_q4_1_cuda(const void * vx, const dfloat * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) { | ||||
|     GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0); | ||||
|     GGML_ASSERT(ncols % (GGML_CUDA_DMMV_X*2) == 0); | ||||
|     const int block_num_y = (nrows + GGML_CUDA_MMV_Y - 1) / GGML_CUDA_MMV_Y; | ||||
|     const dim3 block_nums(block_num_y, 1, 1); | ||||
|     const dim3 block_dims(WARP_SIZE, GGML_CUDA_MMV_Y, 1); | ||||
| @@ -519,7 +519,7 @@ static void dequantize_mul_mat_vec_q4_1_cuda(const void * vx, const dfloat * y, | ||||
| } | ||||
|  | ||||
| static void dequantize_mul_mat_vec_q5_0_cuda(const void * vx, const dfloat * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) { | ||||
|     GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0); | ||||
|     GGML_ASSERT(ncols % (GGML_CUDA_DMMV_X*2) == 0); | ||||
|     const int block_num_y = (nrows + GGML_CUDA_MMV_Y - 1) / GGML_CUDA_MMV_Y; | ||||
|     const dim3 block_nums(block_num_y, 1, 1); | ||||
|     const dim3 block_dims(WARP_SIZE, GGML_CUDA_MMV_Y, 1); | ||||
| @@ -528,7 +528,7 @@ static void dequantize_mul_mat_vec_q5_0_cuda(const void * vx, const dfloat * y, | ||||
| } | ||||
|  | ||||
| static void dequantize_mul_mat_vec_q5_1_cuda(const void * vx, const dfloat * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) { | ||||
|     GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0); | ||||
|     GGML_ASSERT(ncols % (GGML_CUDA_DMMV_X*2) == 0); | ||||
|     const int block_num_y = (nrows + GGML_CUDA_MMV_Y - 1) / GGML_CUDA_MMV_Y; | ||||
|     const dim3 block_nums(block_num_y, 1, 1); | ||||
|     const dim3 block_dims(WARP_SIZE, GGML_CUDA_MMV_Y, 1); | ||||
| @@ -537,7 +537,7 @@ static void dequantize_mul_mat_vec_q5_1_cuda(const void * vx, const dfloat * y, | ||||
| } | ||||
|  | ||||
| static void dequantize_mul_mat_vec_q8_0_cuda(const void * vx, const dfloat * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) { | ||||
|     GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0); | ||||
|     GGML_ASSERT(ncols % (GGML_CUDA_DMMV_X*2) == 0); | ||||
|     const int block_num_y = (nrows + GGML_CUDA_MMV_Y - 1) / GGML_CUDA_MMV_Y; | ||||
|     const dim3 block_nums(block_num_y, 1, 1); | ||||
|     const dim3 block_dims(WARP_SIZE, GGML_CUDA_MMV_Y, 1); | ||||
| @@ -588,7 +588,7 @@ static void dequantize_mul_mat_vec_q6_K_cuda(const void * vx, const float * y, f | ||||
| } | ||||
|  | ||||
| static void convert_mul_mat_vec_f16_cuda(const void * vx, const dfloat * y, float * dst, const int ncols, const int nrows, cudaStream_t stream) { | ||||
|     GGML_ASSERT(ncols % GGML_CUDA_DMMV_X == 0); | ||||
|     GGML_ASSERT(ncols % (GGML_CUDA_DMMV_X*2) == 0); | ||||
|     const int block_num_y = (nrows + GGML_CUDA_MMV_Y - 1) / GGML_CUDA_MMV_Y; | ||||
|     const dim3 block_nums(block_num_y, 1, 1); | ||||
|     const dim3 block_dims(WARP_SIZE, GGML_CUDA_MMV_Y, 1); | ||||
| @@ -672,3 +672,12 @@ void ggml_cuda_op_dequantize_mul_mat_vec( | ||||
|     GGML_UNUSED(src1_ncols); | ||||
|     GGML_UNUSED(src1_padded_row_size); | ||||
| } | ||||
|  | ||||
| bool ggml_cuda_dmmv_type_supported(ggml_type src0_type) { | ||||
|     return src0_type == GGML_TYPE_Q4_0 || src0_type == GGML_TYPE_Q4_1 || | ||||
|         src0_type == GGML_TYPE_Q5_0 || src0_type == GGML_TYPE_Q5_1 || | ||||
|         src0_type == GGML_TYPE_Q8_0 || src0_type == GGML_TYPE_Q2_K || | ||||
|         src0_type == GGML_TYPE_Q3_K || src0_type == GGML_TYPE_Q4_K || | ||||
|         src0_type == GGML_TYPE_Q5_K || src0_type == GGML_TYPE_Q6_K || | ||||
|         src0_type == GGML_TYPE_F16; | ||||
| } | ||||
|   | ||||
| @@ -16,3 +16,5 @@ void ggml_cuda_op_dequantize_mul_mat_vec( | ||||
|     const ggml_tensor * src0, const ggml_tensor * src1, ggml_tensor * dst, const char * src0_dd_i, const float * src1_ddf_i, | ||||
|     const char * src1_ddq_i, float * dst_dd_i, const int64_t row_low, const int64_t row_high, const int64_t src1_ncols, | ||||
|     const int64_t src1_padded_row_size, cudaStream_t stream); | ||||
|  | ||||
| bool ggml_cuda_dmmv_type_supported(ggml_type src0_type); | ||||
|   | ||||
| @@ -804,8 +804,7 @@ struct test_cpy : public test_case { | ||||
|  | ||||
|     test_cpy(ggml_type type_src = GGML_TYPE_F32, ggml_type type_dst = GGML_TYPE_F32, | ||||
|             std::array<int64_t, 4> ne = {10, 10, 10, 1}, | ||||
|             std::array<int64_t, 4> permute = {0, 0, 0, 0}, | ||||
|             bool _dst_use_permute = false) | ||||
|             std::array<int64_t, 4> permute = {0, 0, 0, 0}) | ||||
|         : type_src(type_src), type_dst(type_dst), ne(ne), permute(permute), | ||||
|           _src_use_permute(permute[0] + permute[1] + permute[2] + permute[3] > 0) {} | ||||
|  | ||||
| @@ -2269,6 +2268,8 @@ static bool test_backend(ggml_backend_t backend, test_mode mode, const char * op | ||||
|  | ||||
|     for (ggml_type type_a : other_types) { | ||||
|         for (ggml_type type_b : {GGML_TYPE_F32}) { | ||||
|  | ||||
|             test_cases.emplace_back(new test_mul_mat(type_a, type_b, 16, 1, ggml_blck_size(type_a), { 1,  1}, {1, 1})); | ||||
|             test_cases.emplace_back(new test_mul_mat(type_a, type_b, 16, 1, 256, { 1,  1}, {1, 1})); | ||||
|         } | ||||
|     } | ||||
|   | ||||
		Reference in New Issue
	
	Block a user
	 slaren
					slaren